forked from pytorch/pytorch
-
Notifications
You must be signed in to change notification settings - Fork 1
/
ReplicationPadding.cu
838 lines (715 loc) · 26 KB
/
ReplicationPadding.cu
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
502
503
504
505
506
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
534
535
536
537
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
576
577
578
579
580
581
582
583
584
585
586
587
588
589
590
591
592
593
594
595
596
597
598
599
600
601
602
603
604
605
606
607
608
609
610
611
612
613
614
615
616
617
618
619
620
621
622
623
624
625
626
627
628
629
630
631
632
633
634
635
636
637
638
639
640
641
642
643
644
645
646
647
648
649
650
651
652
653
654
655
656
657
658
659
660
661
662
663
664
665
666
667
668
669
670
671
672
673
674
675
676
677
678
679
680
681
682
683
684
685
686
687
688
689
690
691
692
693
694
695
696
697
698
699
700
701
702
703
704
705
706
707
708
709
710
711
712
713
714
715
716
717
718
719
720
721
722
723
724
725
726
727
728
729
730
731
732
733
734
735
736
737
738
739
740
741
742
743
744
745
746
747
748
749
750
751
752
753
754
755
756
757
758
759
760
761
762
763
764
765
766
767
768
769
770
771
772
773
774
775
776
777
778
779
780
781
782
783
784
785
786
787
788
789
790
791
792
793
794
795
796
797
798
799
800
801
802
803
804
805
806
807
808
809
810
811
812
813
814
815
816
817
818
819
820
821
822
823
824
825
826
827
828
829
830
831
832
833
834
835
836
837
838
#include "ATen/ATen.h"
#include "ATen/cuda/CUDAApplyUtils.cuh"
#include "ATen/cuda/CUDAContext.h"
#include "ATen/NativeFunctions.h"
#include "ATen/TensorUtils.h"
#include "ATen/Utils.h"
#include "c10/util/Exception.h"
#include <THC/THCGeneral.h>
#include "THC/THCNumerics.cuh"
#include "THC/THCDeviceUtils.cuh"
#include <algorithm>
#include <cfloat>
#include <cmath>
namespace at {
namespace native {
__host__ __device__ __forceinline__ int imin(int a, int b) {
return a > b ? b : a;
}
__host__ __device__ __forceinline__ int imax(int a, int b) {
return a > b ? a : b;
}
namespace {
template <typename scalar_t>
__global__ void replication_pad_forward_kernel1d(
PackedTensorAccessor64<scalar_t, 3> input,
PackedTensorAccessor64<scalar_t, 3> output,
int padL, int padR) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= output.size(2)) {
return;
}
int outputPointX = outputPointId % output.size(2);
int iStartX = imax(0, -padL);
int oStartX = imax(0, padL);
int inputPointX = imin(imax(padL, outputPointX), input.size(2) + padL - 1) - oStartX + iStartX;
scalar_t valueToCopy = input[batch][plane][inputPointX];
output[batch][plane][outputPointX] = valueToCopy;
}
template <typename scalar_t>
__global__ void replication_pad_backward_kernel(
PackedTensorAccessor64<scalar_t, 3> gradInput,
PackedTensorAccessor64<scalar_t, 3> gradOutput,
int padL, int padR) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= gradOutput.size(2)) {
return;
}
int outputPointX = outputPointId % gradOutput.size(2);
int iStartX = imax(0, -padL);
int oStartX = imax(0, padL);
int inputPointX = imin(imax(padL, outputPointX), gradInput.size(2) + padL - 1) - oStartX + iStartX;
scalar_t valueToCopy = gradOutput[batch][plane][outputPointX];
atomicAdd(&gradInput[batch][plane][inputPointX], valueToCopy);
}
template <typename scalar_t>
__global__ void replication_pad_forward_kernel2d(
PackedTensorAccessor64<scalar_t, 4> input,
PackedTensorAccessor64<scalar_t, 4> output,
int padT, int padB, int padL, int padR) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= output.size(2) * output.size(3)) {
return;
}
int outputPointX = outputPointId % output.size(3);
int outputPointY = outputPointId / output.size(3);
int iStartX = imax(0, -padL);
int iStartY = imax(0, -padT);
int oStartX = imax(0, padL);
int oStartY = imax(0, padT);
int inputPointX = imin(imax(padL, outputPointX), input.size(3) + padL - 1) - oStartX + iStartX;
int inputPointY = imin(imax(padT, outputPointY), input.size(2) + padT - 1) - oStartY + iStartY;
scalar_t valueToCopy = input[batch][plane][inputPointY][inputPointX];
output[batch][plane][outputPointY][outputPointX] = valueToCopy;
}
template <typename scalar_t>
__global__ void replication_pad_backward_kernel(
PackedTensorAccessor64<scalar_t, 4> gradInput,
PackedTensorAccessor64<scalar_t, 4> gradOutput,
int padT, int padB, int padL, int padR) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= gradOutput.size(2) * gradOutput.size(3)) {
return;
}
int outputPointX = outputPointId % gradOutput.size(3);
int outputPointY = outputPointId / gradOutput.size(3);
int iStartX = imax(0, -padL);
int iStartY = imax(0, -padT);
int oStartX = imax(0, padL);
int oStartY = imax(0, padT);
int inputPointX = imin(imax(padL, outputPointX), gradInput.size(3) + padL - 1) - oStartX + iStartX;
int inputPointY = imin(imax(padT, outputPointY), gradInput.size(2) + padT - 1) - oStartY + iStartY;
scalar_t valueToCopy = gradOutput[batch][plane][outputPointY][outputPointX];
atomicAdd(&gradInput[batch][plane][inputPointY][inputPointX], valueToCopy);
}
template <typename scalar_t>
__global__ void replication_pad_forward_kernel3d(
PackedTensorAccessor64<scalar_t, 5> input,
PackedTensorAccessor64<scalar_t, 5> output,
int pfront, int pback, int ptop, int pbottom, int pleft, int pright) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= (output.size(2) * output.size(3) *
output.size(4))) {
return;
}
int outputPointX = outputPointId % output.size(4);
int outputPointY = (outputPointId / output.size(4)) % output.size(3);
int outputPointZ = outputPointId / (output.size(3) * output.size(4));
int iStartX = imax(0, -pleft);
int iStartY = imax(0, -ptop);
int iStartZ = imax(0, -pfront);
int oStartX = imax(0, pleft);
int oStartY = imax(0, ptop);
int oStartZ = imax(0, pfront);
int inputPointX = imin(imax(pleft, outputPointX),
input.size(4) + pleft - 1) - oStartX + iStartX;
int inputPointY = imin(imax(ptop, outputPointY),
input.size(3) + ptop - 1) - oStartY + iStartY;
int inputPointZ = imin(imax(pfront, outputPointZ),
input.size(2) + pfront - 1) - oStartZ + iStartZ;
scalar_t valueToCopy =
input[batch][plane][inputPointZ][inputPointY][inputPointX];
output[batch][plane][outputPointZ][outputPointY][outputPointX] = valueToCopy;
}
template <typename scalar_t>
__global__ void replication_pad_backward_kernel(
PackedTensorAccessor64<scalar_t, 5> gradInput,
PackedTensorAccessor64<scalar_t, 5> gradOutput,
int pfront, int pback, int ptop, int pbottom, int pleft, int pright) {
int outputPointId = threadIdx.x + blockIdx.x * blockDim.x;
int plane = blockIdx.y;
int batch = blockIdx.z;
if (outputPointId >= (gradOutput.size(2) * gradOutput.size(3) *
gradOutput.size(4))) {
return;
}
int outputPointX = outputPointId % gradOutput.size(4);
int outputPointY = (outputPointId / gradOutput.size(4)) %
gradOutput.size(3);
int outputPointZ = outputPointId / (gradOutput.size(3) *
gradOutput.size(4));
int iStartX = imax(0, -pleft);
int iStartY = imax(0, -ptop);
int iStartZ = imax(0, -pfront);
int oStartX = imax(0, pleft);
int oStartY = imax(0, ptop);
int oStartZ = imax(0, pfront);
int inputPointX = imin(imax(pleft, outputPointX),
gradInput.size(4) + pleft - 1) - oStartX + iStartX;
int inputPointY = imin(imax(ptop, outputPointY),
gradInput.size(3) + ptop - 1) - oStartY + iStartY;
int inputPointZ = imin(imax(pfront, outputPointZ),
gradInput.size(2) + pfront - 1) - oStartZ + iStartZ;
scalar_t valueToCopy =
gradOutput[batch][plane][outputPointZ][outputPointY][outputPointX];
atomicAdd(&gradInput[batch][plane][inputPointZ][inputPointY][inputPointX],
valueToCopy);
}
void replication_pad1d_out_cuda_template(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
TORCH_CHECK(paddingSize.size() == 2, "padding Size is expected to be 2");
int padL = paddingSize[0];
int padR = paddingSize[1];
int planeDim = 0;
int dimw = 1;
int numBatch = 1;
int numInputDims = input.ndimension();
TORCH_CHECK(input.numel() > 0 && (numInputDims == 2 || numInputDims == 3),
"2D or 3D (batch mode) tensor expected for input")
if (numInputDims == 3) {
numBatch = input.size(0);
planeDim++;
dimw++;
}
int numPlanes = input.size(planeDim);
int inputW = input.size(dimw);
int outputW = inputW + padL + padR;
TORCH_CHECK(outputW >= 1,
"input (W: ", inputW, ")is too small."
" Calculated output W: ", outputW);
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad1d_cuda", [&] {
if (numInputDims == 2) {
output.resize_({numPlanes, outputW});
auto input_ = input.unsqueeze(0);
auto output_ = output.unsqueeze(0);
auto devInput = input_.packed_accessor64<scalar_t, 3>();
auto devOutput = output_.packed_accessor64<scalar_t, 3>();
int outputPlaneSize = devOutput.size(2);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel1d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(devInput, devOutput, padL, padR);
} else {
output.resize_({numBatch, numPlanes, outputW});
auto devInput = input.packed_accessor64<scalar_t, 3>();
auto devOutput = output.packed_accessor64<scalar_t, 3>();
int outputPlaneSize = devOutput.size(2);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel1d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(devInput, devOutput, padL, padR);
}
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
void replication_pad1d_backward_out_cuda_template(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(gradOutput),
"output gradient tensor must fit into 32-bit index math");
TORCH_CHECK(paddingSize.size() == 2, "padding Size is expected to be 2");
int padL = paddingSize[0];
int padR = paddingSize[1];
int planeDim = 0;
int dimw = 1;
int numInputDims = input.ndimension();
if (numInputDims == 3) {
planeDim++;
dimw++;
}
int iwidth = input.size(dimw);
int owidth = iwidth + padL + padR;
TORCH_CHECK(owidth == gradOutput.size(dimw),
"gradOutput width unexpected. Expected: ", owidth, ", Got: ",
gradOutput.size(dimw));
gradInput.resize_as_(input);
gradInput.zero_();
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad1d_backward_cuda", [&] {
auto gradInput_ = gradInput;
auto gradOutput_ = gradOutput;
if (numInputDims == 2) {
gradInput_ = gradInput.unsqueeze(0);
gradOutput_ = gradOutput.unsqueeze(0);
}
auto devGradInput = gradInput_.packed_accessor64<scalar_t, 3>();
auto devGradOutput = gradOutput_.packed_accessor64<scalar_t, 3>();
int outputPlaneSize = devGradOutput.size(2);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devGradOutput.size(1),
devGradOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_backward_kernel <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(devGradInput, devGradOutput,
padL, padR);
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
void replication_pad2d_out_cuda_template(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
TORCH_CHECK(paddingSize.size() == 4, "padding Size is expected to be 4");
int padL = paddingSize[0];
int padR = paddingSize[1];
int padT = paddingSize[2];
int padB = paddingSize[3];
int planeDim = 0;
int dimh = 1;
int dimw = 2;
int numBatch = 1;
int numInputDims = input.dim();
TORCH_CHECK(input.numel() && (numInputDims == 3 || numInputDims == 4),
"non-empty 3D or 4D (batch mode) tensor expected for input, but got: ",
input)
if (numInputDims == 4) {
numBatch = input.size(0);
planeDim++;
dimh++;
dimw++;
}
int numPlanes = input.size(planeDim);
int inputH = input.size(dimh);
int inputW = input.size(dimw);
int outputH = inputH + padT + padB;
int outputW = inputW + padL + padR;
TORCH_CHECK(outputW >= 1 || outputH >= 1,
"input (H: ", inputH, ", W: ", inputW, ") is too small."
" Calculated output H: ", outputH, " W: ", outputW);
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad2d_cuda", [&] {
if (numInputDims == 3) {
output.resize_({numPlanes, outputH, outputW});
auto input_ = input.unsqueeze(0);
auto output_ = output.unsqueeze(0);
auto devInput = input_.packed_accessor64<scalar_t, 4>();
auto devOutput = output_.packed_accessor64<scalar_t, 4>();
int outputPlaneSize = devOutput.size(2) * devOutput.size(3);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel2d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(
devInput, devOutput, padT, padB, padL, padR);
} else {
output.resize_({numBatch, numPlanes, outputH, outputW});
auto devInput = input.packed_accessor64<scalar_t, 4>();
auto devOutput = output.packed_accessor64<scalar_t, 4>();
int outputPlaneSize = devOutput.size(2) * devOutput.size(3);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel2d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(devInput, devOutput,
padT, padB, padL, padR);
}
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
void replication_pad2d_backward_out_cuda_template(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(gradOutput),
"output gradient tensor must fit into 32-bit index math");
TORCH_CHECK(paddingSize.size() == 4, "padding Size is expected to be 4");
int padL = paddingSize[0];
int padR = paddingSize[1];
int padT = paddingSize[2];
int padB = paddingSize[3];
int planeDim = 0;
int dimh = 1;
int dimw = 2;
int numInputDims = input.dim();
if (numInputDims == 4) {
planeDim++;
dimh++;
dimw++;
}
int iheight = input.size(dimh);
int iwidth = input.size(dimw);
int oheight = iheight + padT + padB;
int owidth = iwidth + padL + padR;
TORCH_CHECK(owidth == gradOutput.size(dimw),
"gradOutput width unexpected. Expected: ", owidth, ", Got: ",
gradOutput.size(dimw));
TORCH_CHECK(oheight == gradOutput.size(dimh),
"gradOutput height unexpected. Expected: ", oheight, ", Got: ",
gradOutput.size(dimh));
gradInput.resize_as_(input);
gradInput.zero_();
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad2d_backward_cuda", [&] {
auto gradInput_ = gradInput;
auto gradOutput_ = gradOutput;
if (numInputDims == 3) {
gradInput_ = gradInput.unsqueeze(0);
gradOutput_ = gradOutput.unsqueeze(0);
}
auto devGradInput = gradInput_.packed_accessor64<scalar_t, 4>();
auto devGradOutput = gradOutput_.packed_accessor64<scalar_t, 4>();
int outputPlaneSize = devGradOutput.size(2) * devGradOutput.size(3);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devGradOutput.size(1),
devGradOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_backward_kernel <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(devGradInput, devGradOutput,
padT, padB, padL, padR);
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
static inline void shapeCheck3d(
const Tensor& input,
int pleft, int pright,
int ptop, int pbottom,
int pfront, int pback) {
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
int numInputDims = input.dim();
TORCH_CHECK(input.numel() && (numInputDims == 4 || numInputDims == 5),
"non-empty 4D or 5D (batch mode) tensor expected for input, but got: ", input);
int planeDim = 0;
int dimd = 1;
int dimh = 2;
int dimw = 3;
if (numInputDims == 5) {
planeDim++;
dimd++;
dimh++;
dimw++;
}
int numPlanes = input.size(planeDim);
int idepth = input.size(dimd);
int iheight = input.size(dimh);
int iwidth = input.size(dimw);
int odepth = idepth + pfront + pback;
int oheight = iheight + ptop + pbottom;
int owidth = iwidth + pleft + pright;
TORCH_CHECK(owidth >= 1 || oheight >= 1 || odepth >= 1,
"input (D: ", idepth, " H: ", iheight, ", W: ", iwidth,
") is too small."
" Calculated output D: ", odepth, " H: ", oheight, " W: ", owidth);
}
static inline void shapeAndGradOutputCheck3d(
const Tensor& input,
const Tensor& gradOutput,
int pleft, int pright,
int ptop, int pbottom,
int pfront, int pback) {
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(input),
"input tensor must fit into 32-bit index math");
int numInputDims = input.dim();
TORCH_CHECK(input.numel() && (numInputDims == 4 || numInputDims == 5),
"non-empty 4D or 5D (batch mode) tensor expected for input, but got: ", input);
int planeDim = 0;
int dimd = 1;
int dimh = 2;
int dimw = 3;
if (numInputDims == 5) {
planeDim++;
dimd++;
dimh++;
dimw++;
}
int numPlanes = input.size(planeDim);
int idepth = input.size(dimd);
int iheight = input.size(dimh);
int iwidth = input.size(dimw);
int odepth = idepth + pfront + pback;
int oheight = iheight + ptop + pbottom;
int owidth = iwidth + pleft + pright;
TORCH_CHECK(owidth >= 1 || oheight >= 1 || odepth >= 1,
"input (D: ", idepth, " H: ", iheight, ", W: ", iwidth,
") is too small."
" Calculated output D: ", odepth, " H: ", oheight, " W: ", owidth);
TORCH_CHECK(at::cuda::detail::canUse32BitIndexMath(gradOutput),
"output gradient tensor must fit into 32-bit index math");
TORCH_CHECK(numPlanes == gradOutput.size(planeDim),
"gradOutput width unexpected. Expected: ", numPlanes, ", Got: ",
gradOutput.size(planeDim));
TORCH_CHECK(owidth == gradOutput.size(dimw),
"gradOutput width unexpected. Expected: ", owidth, ", Got: ",
gradOutput.size(dimw));
TORCH_CHECK(oheight == gradOutput.size(dimh),
"gradOutput height unexpected. Expected: ", oheight, ", Got: ",
gradOutput.size(dimh));
TORCH_CHECK(odepth == gradOutput.size(dimd),
"gradOutput depth unexpected. Expected: ", odepth, ", Got: ",
gradOutput.size(dimd));
}
void replication_pad3d_out_cuda_template(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(paddingSize.size() == 6, "padding Size is expected to be 6");
int pleft = paddingSize[0];
int pright = paddingSize[1];
int ptop = paddingSize[2];
int pbottom = paddingSize[3];
int pfront = paddingSize[4];
int pback = paddingSize[5];
shapeCheck3d(input, pleft, pright, ptop,
pbottom, pfront, pback);
int planeDim = 0;
int dimd = 1;
int dimh = 2;
int dimw = 3;
int numBatch = 1;
int numInputDims = input.dim();
if (numInputDims == 5) {
numBatch = input.size(0);
planeDim++;
dimd++;
dimh++;
dimw++;
}
int numPlanes = input.size(planeDim);
int inputD = input.size(dimd);
int inputH = input.size(dimh);
int inputW = input.size(dimw);
int outputD = inputD + pfront + pback;
int outputH = inputH + ptop + pbottom;
int outputW = inputW + pleft + pright;
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad3d_cuda", [&] {
if (numInputDims == 4) {
output.resize_({numPlanes, outputD, outputH, outputW});
auto input_ = input.unsqueeze(0);
auto output_ = output.unsqueeze(0);
auto devInput = input_.packed_accessor64<scalar_t, 5>();
auto devOutput = output_.packed_accessor64<scalar_t, 5>();
int outputPlaneSize = devOutput.size(2) * devOutput.size(3) *
devOutput.size(4);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel3d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(
devInput, devOutput, pfront, pback, ptop, pbottom, pleft, pright);
} else {
output.resize_({numBatch, numPlanes, outputD, outputH, outputW});
auto devInput = input.packed_accessor64<scalar_t, 5>();
auto devOutput = output.packed_accessor64<scalar_t, 5>();
int outputPlaneSize = devOutput.size(2) * devOutput.size(3) *
devOutput.size(4);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devOutput.size(1),
devOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_forward_kernel3d <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(
devInput, devOutput, pfront, pback, ptop, pbottom, pleft, pright);
}
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
void replication_pad3d_backward_out_cuda_template(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
TORCH_CHECK(paddingSize.size() == 6, "padding Size is expected to be 6");
int pleft = paddingSize[0];
int pright = paddingSize[1];
int ptop = paddingSize[2];
int pbottom = paddingSize[3];
int pfront = paddingSize[4];
int pback = paddingSize[5];
shapeAndGradOutputCheck3d(input, gradOutput, pleft, pright, ptop,
pbottom, pfront, pback);
int planeDim = 0;
int dimd = 1;
int dimh = 2;
int dimw = 3;
int numInputDims = input.dim();
if (numInputDims == 5) {
planeDim++;
dimd++;
dimh++;
dimw++;
}
gradInput.resize_as_(input);
gradInput.zero_();
AT_DISPATCH_FLOATING_TYPES_AND_HALF(
input.scalar_type(), "replication_pad3d_backward_cuda", [&] {
auto gradInput_ = gradInput;
auto gradOutput_ = gradOutput;
if (numInputDims == 4) {
gradInput_ = gradInput.unsqueeze(0);
gradOutput_ = gradOutput.unsqueeze(0);
}
auto devGradInput = gradInput_.packed_accessor64<scalar_t, 5>();
auto devGradOutput = gradOutput_.packed_accessor64<scalar_t, 5>();
int outputPlaneSize = devGradOutput.size(2) * devGradOutput.size(3) *
devGradOutput.size(4);
dim3 gridSize(THCCeilDiv(outputPlaneSize, 256),
devGradOutput.size(1),
devGradOutput.size(0));
dim3 blockSize(outputPlaneSize > 256 ? 256 : outputPlaneSize);
replication_pad_backward_kernel <<<gridSize, blockSize, 0,
at::cuda::getCurrentCUDAStream()>>>(
devGradInput, devGradOutput, pfront, pback, ptop, pbottom, pleft, pright);
}
);
AT_CUDA_CHECK(cudaGetLastError());
}
} // namespace
Tensor& replication_pad1d_out_cuda(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad1d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor replication_pad1d_cuda(
const Tensor& input,
IntArrayRef paddingSize)
{
auto output = at::empty({0}, input.options());
replication_pad1d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor& replication_pad1d_backward_out_cuda(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad1d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
Tensor replication_pad1d_backward_cuda(
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
auto gradInput = at::zeros_like(input, LEGACY_CONTIGUOUS_MEMORY_FORMAT);
replication_pad1d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
Tensor& replication_pad2d_out_cuda(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad2d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor replication_pad2d_cuda(
const Tensor& input,
IntArrayRef paddingSize)
{
auto output = at::empty({0}, input.options());
replication_pad2d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor& replication_pad2d_backward_out_cuda(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad2d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
Tensor replication_pad2d_backward_cuda(
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
auto gradInput = at::zeros_like(input, LEGACY_CONTIGUOUS_MEMORY_FORMAT);
replication_pad2d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
Tensor& replication_pad3d_out_cuda(
Tensor& output,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad3d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor replication_pad3d_cuda(
const Tensor& input,
IntArrayRef paddingSize)
{
auto output = at::empty({0}, input.options());
replication_pad3d_out_cuda_template(
output, input, paddingSize);
return output;
}
Tensor& replication_pad3d_backward_out_cuda(
Tensor& gradInput,
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
replication_pad3d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
Tensor replication_pad3d_backward_cuda(
const Tensor& gradOutput,
const Tensor& input,
IntArrayRef paddingSize)
{
auto gradInput = at::zeros_like(input, LEGACY_CONTIGUOUS_MEMORY_FORMAT);
replication_pad3d_backward_out_cuda_template(
gradInput, gradOutput, input, paddingSize);
return gradInput;
}
} // at::native
} // at