Skip to content

Commit 9775db5

Browse files
authored
[0-size Tensor No.181] Add 0-size Tensor support for paddle.nn.functional.interpolate [fluid_ops]
1 parent 353502b commit 9775db5

File tree

10 files changed

+95
-8
lines changed

10 files changed

+95
-8
lines changed

paddle/phi/infermeta/multiary.cc

Lines changed: 3 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -3245,7 +3245,7 @@ static void Interpolate1DInferShapeCheck(
32453245
"Input(X) dimension is 3, but got method = %s .",
32463246
interp_method));
32473247
const DataLayout data_layout = common::StringToDataLayout(data_layout_str);
3248-
for (int i = 0; i < dim_x.size(); ++i) {
3248+
for (int i = 2; i < dim_x.size(); ++i) {
32493249
PADDLE_ENFORCE_NE(dim_x[i],
32503250
0,
32513251
common::errors::InvalidArgument(
@@ -3377,7 +3377,7 @@ static void Interpolate2DInferShapeCheck(
33773377
interp_method));
33783378
const DataLayout data_layout = common::StringToDataLayout(data_layout_str);
33793379

3380-
for (int i = 0; i < dim_x.size(); ++i) {
3380+
for (int i = 2; i < dim_x.size(); ++i) {
33813381
PADDLE_ENFORCE_NE(dim_x[i],
33823382
0,
33833383
common::errors::InvalidArgument(
@@ -3530,7 +3530,7 @@ static void Interpolate3DInferShapeCheck(
35303530
interp_method));
35313531
const DataLayout data_layout = common::StringToDataLayout(data_layout_str);
35323532

3533-
for (int i = 0; i < dim_x.size(); ++i) {
3533+
for (int i = 2; i < dim_x.size(); ++i) {
35343534
PADDLE_ENFORCE_NE(dim_x[i],
35353535
0,
35363536
common::errors::InvalidArgument(

paddle/phi/kernels/cpu/interpolate_grad_kernel.cc

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -861,6 +861,10 @@ void InterpolateGradKernel(
861861
bool align_corners,
862862
int align_mode,
863863
DenseTensor* x_grad) {
864+
if (x_grad && x_grad->numel() == 0) {
865+
dev_ctx.template Alloc<T>(x_grad);
866+
return;
867+
}
864868
auto output_grad_dims = output_grad.dims();
865869
if (output_grad_dims.size() == 3) { // 1D interpolation grad
866870
Interpolate1DCPUBwd<T, Context>(dev_ctx,

paddle/phi/kernels/cpu/interpolate_kernel.cc

Lines changed: 5 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1033,6 +1033,11 @@ void InterpolateKernel(
10331033
bool align_corners,
10341034
int align_mode,
10351035
DenseTensor* output) {
1036+
if (x.numel() == 0) {
1037+
dev_ctx.template Alloc<T>(output);
1038+
return;
1039+
}
1040+
10361041
auto input_dims = x.dims();
10371042
if (input_dims.size() == 3) { // 1D interpolation
10381043
Interpolate1DCPUFwd<T, Context>(dev_ctx,

paddle/phi/kernels/gpu/interpolate_grad_kernel.cu

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1438,6 +1438,10 @@ void InterpolateGradKernel(
14381438
bool align_corners,
14391439
int align_mode,
14401440
DenseTensor* x_grad) {
1441+
if (x_grad && x_grad->numel() == 0) {
1442+
dev_ctx.template Alloc<T>(x_grad);
1443+
return;
1444+
}
14411445
auto output_grad_dims = output_grad.dims();
14421446
if (output_grad_dims.size() == 3) { // 1D interpolation grad
14431447
Interpolate1DCUDABwd<T, Context>(dev_ctx,

paddle/phi/kernels/gpu/interpolate_kernel.cu

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1230,6 +1230,10 @@ void InterpolateKernel(
12301230
bool align_corners,
12311231
int align_mode,
12321232
DenseTensor* output) {
1233+
if (x.numel() == 0) {
1234+
dev_ctx.template Alloc<T>(output);
1235+
return;
1236+
}
12331237
auto input_dims = x.dims();
12341238
if (input_dims.size() == 3) { // 1D interpolation
12351239
Interpolate1DCUDAFwd<T, Context>(dev_ctx,

paddle/phi/kernels/xpu/interpolate_grad_kernel.cc

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -39,6 +39,10 @@ void InterpolateGradKernel(
3939
bool align_corners,
4040
int align_mode,
4141
DenseTensor* x_grad) {
42+
if (x_grad && x_grad->numel() == 0) {
43+
dev_ctx.template Alloc<T>(x_grad);
44+
return;
45+
}
4246
const DataLayout data_layout = common::StringToDataLayout(data_layout_str);
4347
int n, c, in_d, in_h, in_w;
4448
funcs::ExtractNCDWH(x.dims(), data_layout, &n, &c, &in_d, &in_h, &in_w);

paddle/phi/kernels/xpu/interpolate_kernel.cc

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -38,6 +38,10 @@ void InterpolateKernel(
3838
bool align_corners,
3939
int align_mode,
4040
DenseTensor* output) {
41+
if (x.numel() == 0) {
42+
dev_ctx.template Alloc<T>(output);
43+
return;
44+
}
4145
using XPUType = typename XPUTypeTrait<T>::Type;
4246
const DataLayout data_layout = common::StringToDataLayout(data_layout_str);
4347
int n, c, in_d, in_h, in_w;

python/paddle/nn/functional/common.py

Lines changed: 38 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -642,7 +642,7 @@ def _is_list_or_tuple_(data):
642642
if len(x.shape) == 4:
643643
if len(out_shape) != 2:
644644
raise ValueError(
645-
"size length should be 2 for " "input 4-D tensor."
645+
"size length should be 2 for input 4-D tensor."
646646
)
647647
if contain_var:
648648
attrs['out_h'] = size_list[0]
@@ -667,6 +667,30 @@ def _is_list_or_tuple_(data):
667667
attrs['out_w'] = out_shape[2]
668668

669669
elif scale is not None:
670+
# scale in python is float64, but in kernel is float32, so we need to recalculate the scale in float32
671+
# Currently it is only used when x.size is 0.
672+
x_shape = x.shape
673+
if data_format == 'NCW':
674+
max_dim = x_shape[2]
675+
elif data_format == 'NWC':
676+
max_dim = x_shape[1]
677+
elif data_format == 'NCHW':
678+
max_dim = max(x.shape[2], x.shape[3])
679+
elif data_format == 'NHWC':
680+
max_dim = max(x.shape[1], x.shape[2])
681+
elif data_format == 'NCDHW':
682+
max_dim = max(x.shape[2], x.shape[3], x.shape[4])
683+
elif data_format == 'NDHWC':
684+
max_dim = max(x.shape[1], x.shape[2], x.shape[3])
685+
else:
686+
max_dim = 1
687+
688+
def _scale_to_float32(value):
689+
if len(str(value)) <= 10:
690+
return value
691+
# round down
692+
return numpy.float32(int(value * max_dim) / max_dim)
693+
670694
if recompute_scale_factor:
671695
if in_dynamic_mode() and isinstance(scale, Variable):
672696
if scale.shape == []:
@@ -710,11 +734,15 @@ def _is_list_or_tuple_(data):
710734

711735
scale = None
712736
else:
713-
if in_dynamic_mode() and isinstance(scale, Variable):
737+
dynamic_mode = False
738+
if in_dynamic_mode():
739+
dynamic_mode = True
740+
if dynamic_mode and isinstance(scale, Variable):
714741
if scale.shape == []:
715742
scale = float(scale)
716743
else:
717744
scale = list(scale.numpy())
745+
718746
if isinstance(scale, (Variable, paddle.pir.Value)):
719747
scale.stop_gradient = True
720748
inputs["Scale"] = scale
@@ -724,7 +752,10 @@ def _is_list_or_tuple_(data):
724752
scale_list = []
725753
for i in range(len(x.shape) - 2):
726754
scale_list.append(scale)
727-
attrs['scale'] = list(map(float, scale_list))
755+
if dynamic_mode and x.size == 0:
756+
attrs['scale'] = list(map(_scale_to_float32, scale_list))
757+
else:
758+
attrs['scale'] = list(map(float, scale_list))
728759
elif isinstance(scale, (list, tuple)):
729760
if len(scale) != len(x.shape) - 2:
730761
raise ValueError(
@@ -736,7 +767,10 @@ def _is_list_or_tuple_(data):
736767
raise ValueError(
737768
"Attr(scale) should be greater than zero."
738769
)
739-
attrs['scale'] = list(map(float, scale))
770+
if dynamic_mode and x.size == 0:
771+
attrs['scale'] = list(map(_scale_to_float32, scale))
772+
else:
773+
attrs['scale'] = list(map(float, scale))
740774
else:
741775
raise TypeError(
742776
"Attr(scale)'s type should be float, int, list, tuple, or Tensor."

test/legacy_test/test_adaptive_avg_pool2d.py

Lines changed: 28 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -515,5 +515,33 @@ def test_grad(self):
515515
np.testing.assert_allclose(x.grad.shape, x.shape)
516516

517517

518+
class TestInterpolateAPI_ZeroSize(unittest.TestCase):
519+
def setUp(self):
520+
self.x_np = np.random.random([0, 3, 7, 7]).astype("float32")
521+
522+
def test_functional_interpolate(self):
523+
for use_cuda in (
524+
[False, True] if core.is_compiled_with_cuda() else [False]
525+
):
526+
place = paddle.CUDAPlace(0) if use_cuda else paddle.CPUPlace()
527+
paddle.disable_static(place=place)
528+
x = paddle.to_tensor(self.x_np)
529+
x.stop_gradient = False
530+
531+
out = paddle.nn.functional.interpolate(
532+
x=x, mode="area", size=[2, 5]
533+
)
534+
res_np = adaptive_pool2d_forward(
535+
x=self.x_np, output_size=[2, 5], pool_type="avg"
536+
)
537+
np.testing.assert_allclose(
538+
out.numpy(), res_np, rtol=1e-5, atol=1e-8
539+
)
540+
541+
loss = paddle.sum(out)
542+
loss.backward()
543+
np.testing.assert_allclose(x.grad.shape, x.shape)
544+
545+
518546
if __name__ == '__main__':
519547
unittest.main()

test/legacy_test/test_bicubic_interp_v2_op.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -912,7 +912,7 @@ def test_input_shape_1():
912912
self.assertRaises(ValueError, test_size_length)
913913
self.assertRaises(ValueError, test_size_tensor_ndim)
914914
self.assertRaises(ValueError, test_size_tensor_length)
915-
self.assertRaises(ValueError, test_input_shape_1)
915+
# self.assertRaises(ValueError, test_input_shape_1)
916916

917917
def test_errors(self):
918918
with program_guard(Program(), Program()):

0 commit comments

Comments
 (0)