Skip to content

Commit a1749c6

Browse files
committed
update unpool
1 parent 5322bd0 commit a1749c6

File tree

8 files changed

+158
-17
lines changed

8 files changed

+158
-17
lines changed

paddle/phi/kernels/cpu/unpool_grad_kernel.cc

Lines changed: 14 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -130,8 +130,18 @@ void Unpool3dGradKernel(const Context& dev_ctx,
130130

131131
} // namespace phi
132132

133-
PD_REGISTER_KERNEL(
134-
unpool_grad, CPU, ALL_LAYOUT, phi::UnpoolGradKernel, float, double) {}
133+
PD_REGISTER_KERNEL(unpool_grad,
134+
CPU,
135+
ALL_LAYOUT,
136+
phi::UnpoolGradKernel,
137+
float,
138+
double,
139+
int64_t) {}
135140

136-
PD_REGISTER_KERNEL(
137-
unpool3d_grad, CPU, ALL_LAYOUT, phi::Unpool3dGradKernel, float, double) {}
141+
PD_REGISTER_KERNEL(unpool3d_grad,
142+
CPU,
143+
ALL_LAYOUT,
144+
phi::Unpool3dGradKernel,
145+
float,
146+
double,
147+
int64_t) {}

paddle/phi/kernels/cpu/unpool_kernel.cc

Lines changed: 3 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -126,7 +126,8 @@ void Unpool3dKernel(const Context& dev_ctx,
126126

127127
} // namespace phi
128128

129-
PD_REGISTER_KERNEL(unpool, CPU, ALL_LAYOUT, phi::UnpoolKernel, float, double) {}
129+
PD_REGISTER_KERNEL(
130+
unpool, CPU, ALL_LAYOUT, phi::UnpoolKernel, float, double, int64_t) {}
130131

131132
PD_REGISTER_KERNEL(
132-
unpool3d, CPU, ALL_LAYOUT, phi::Unpool3dKernel, float, double) {}
133+
unpool3d, CPU, ALL_LAYOUT, phi::Unpool3dKernel, float, double, int64_t) {}

paddle/phi/kernels/gpu/unpool_grad_kernel.cu

Lines changed: 14 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -188,8 +188,18 @@ void Unpool3dGradKernel(const Context& dev_ctx,
188188

189189
} // namespace phi
190190

191-
PD_REGISTER_KERNEL(
192-
unpool_grad, GPU, ALL_LAYOUT, phi::UnpoolGradKernel, float, double) {}
191+
PD_REGISTER_KERNEL(unpool_grad,
192+
GPU,
193+
ALL_LAYOUT,
194+
phi::UnpoolGradKernel,
195+
float,
196+
double,
197+
int64_t) {}
193198

194-
PD_REGISTER_KERNEL(
195-
unpool3d_grad, GPU, ALL_LAYOUT, phi::Unpool3dGradKernel, float, double) {}
199+
PD_REGISTER_KERNEL(unpool3d_grad,
200+
GPU,
201+
ALL_LAYOUT,
202+
phi::Unpool3dGradKernel,
203+
float,
204+
double,
205+
int64_t) {}

paddle/phi/kernels/gpu/unpool_kernel.cu

Lines changed: 9 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -173,7 +173,13 @@ void Unpool3dKernel(const Context& dev_ctx,
173173
} // namespace phi
174174

175175
PD_REGISTER_KERNEL(
176-
unpool, GPU, ALL_LAYOUT, phi::UnpoolKernel, int, float, double) {}
176+
unpool, GPU, ALL_LAYOUT, phi::UnpoolKernel, int, float, double, int64_t) {}
177177

178-
PD_REGISTER_KERNEL(
179-
unpool3d, GPU, ALL_LAYOUT, phi::Unpool3dKernel, int, float, double) {}
178+
PD_REGISTER_KERNEL(unpool3d,
179+
GPU,
180+
ALL_LAYOUT,
181+
phi::Unpool3dKernel,
182+
int,
183+
float,
184+
double,
185+
int64_t) {}

python/paddle/nn/functional/pooling.py

Lines changed: 6 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -755,7 +755,7 @@ def max_unpool1d(
755755
x (Tensor): The input tensor of unpooling operator which is a 3-D tensor with
756756
shape [N, C, L]. The format of input tensor is `"NCL"`,
757757
where `N` is batch size, `C` is the number of channels, `L` is
758-
the length of the feature. The data type is float32 or float64.
758+
the length of the feature. The data type is float32, float64 or int64.
759759
indices (Tensor): The indices given out by maxpooling1d which is a 3-D tensor with
760760
shape [N, C, L]. The format of input tensor is `"NCL"` ,
761761
where `N` is batch size, `C` is the number of channels, `L` is
@@ -813,6 +813,8 @@ def max_unpool1d(
813813
# use 2d to implenment 1d should expand padding in advance.
814814
padding = _expand_low_nd_padding(padding)
815815

816+
if output_size is not None:
817+
output_size = output_size[:2] + [1] + output_size[2:]
816818
output_size = _unpool_output_size(
817819
x, kernel_size, stride, padding, output_size
818820
)
@@ -863,12 +865,12 @@ def max_unpool2d(
863865
shape [N, C, H, W]. The format of input tensor is `"NCHW"`,
864866
where `N` is batch size, `C` is the number of channels,
865867
`H` is the height of the feature, and `W` is the width of the
866-
feature. The data type if float32 or float64.
868+
feature. The data type is float32, float64 or int64.
867869
indices (Tensor): The indices given out by maxpooling2d which is a 4-D tensor with
868870
shape [N, C, H, W]. The format of input tensor is `"NCHW"` ,
869871
where `N` is batch size, `C` is the number of channels,
870872
`H` is the height of the feature, and `W` is the width of the
871-
feature. The data type if float32 or float64.
873+
feature. The data type is float32 or float64.
872874
kernel_size (int|list|tuple): The unpool kernel size. If unpool kernel size is a tuple or list,
873875
it must contain an integer.
874876
stride (int|list|tuple): The unpool stride size. If unpool stride size is a tuple or list,
@@ -1011,7 +1013,7 @@ def max_unpool3d(
10111013
shape [N, C, D, H, W]. The format of input tensor is `"NCDHW"`,
10121014
where `N` is batch size, `C` is the number of channels, `D` is
10131015
the depth of the feature, `H` is the height of the feature,
1014-
and `W` is the width of the feature. The data type is float32 or float64.
1016+
and `W` is the width of the feature. The data type is float32, float64 or int64.
10151017
indices (Tensor): The indices given out by maxpooling3d which is a 5-D tensor with
10161018
shape [N, C, D, H, W]. The format of input tensor is `"NCDHW"` ,
10171019
where `N` is batch size, `C` is the number of channels, `D` is

test/legacy_test/test_unpool1d_op.py

Lines changed: 29 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -135,6 +135,35 @@ def test_case(self):
135135
paddle.enable_static()
136136

137137

138+
class TestUnpool1DOpAPI_dygraph4(unittest.TestCase):
139+
def test_case(self):
140+
places = [paddle.CPUPlace()]
141+
if paddle.base.core.is_compiled_with_cuda():
142+
places.append(paddle.CUDAPlace(0))
143+
for place in places:
144+
paddle.disable_static()
145+
input_data = np.arange(3 * 16).reshape([1, 3, 16]).astype("float32")
146+
input_x = paddle.to_tensor(input_data)
147+
output, indices = F.max_pool1d(
148+
input_x, kernel_size=2, stride=2, return_mask=True
149+
)
150+
output_unpool = F.max_unpool1d(
151+
output.astype("int64"),
152+
indices,
153+
kernel_size=2,
154+
stride=2,
155+
output_size=input_x.shape,
156+
)
157+
expected_output_unpool = unpool1dmax_forward_naive(
158+
output.numpy(), indices.numpy(), [2], [2], [0], [16]
159+
)
160+
np.testing.assert_allclose(
161+
output_unpool.numpy(), expected_output_unpool, rtol=1e-05
162+
)
163+
164+
paddle.enable_static()
165+
166+
138167
class TestUnpool1DOpAPI_static(unittest.TestCase):
139168
@test_with_pir_api
140169
def test_case(self):

test/legacy_test/test_unpool3d_op.py

Lines changed: 38 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -373,6 +373,44 @@ def test_case(self):
373373
paddle.enable_static()
374374

375375

376+
class TestUnpool3DOpAPI_dygraph4(unittest.TestCase):
377+
def test_case(self):
378+
places = [paddle.CPUPlace()]
379+
if paddle.base.core.is_compiled_with_cuda():
380+
places.append(paddle.CUDAPlace(0))
381+
for place in places:
382+
paddle.disable_static()
383+
input_data = (
384+
np.arange(3 * 4 * 4 * 6)
385+
.reshape([1, 3, 4, 4, 6])
386+
.astype("float32")
387+
)
388+
input_x = paddle.to_tensor(input_data)
389+
output, indices = F.max_pool3d(
390+
input_x, kernel_size=2, stride=2, return_mask=True
391+
)
392+
output_unpool = F.max_unpool3d(
393+
output.astype("int64"),
394+
indices,
395+
kernel_size=2,
396+
stride=2,
397+
output_size=input_x.shape,
398+
)
399+
expected_output_unpool = unpool3dmax_forward_naive(
400+
output.numpy(),
401+
indices.numpy(),
402+
[2, 2, 2],
403+
[2, 2, 2],
404+
[0, 0, 0],
405+
[4, 4, 6],
406+
)
407+
np.testing.assert_allclose(
408+
output_unpool.numpy(), expected_output_unpool, rtol=1e-05
409+
)
410+
411+
paddle.enable_static()
412+
413+
376414
class TestUnpool3DOpAPI_static(unittest.TestCase):
377415
@test_with_pir_api
378416
def test_case(self):

test/legacy_test/test_unpool_op.py

Lines changed: 45 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -400,6 +400,51 @@ def test_case(self):
400400
np.testing.assert_allclose(out_pp.numpy(), expect_res, rtol=1e-05)
401401

402402

403+
class TestUnpoolOpAPI_dy4(unittest.TestCase):
404+
def test_case(self):
405+
import numpy as np
406+
407+
import paddle
408+
import paddle.nn.functional as F
409+
from paddle import base
410+
from paddle.base import core
411+
412+
if core.is_compiled_with_cuda():
413+
place = core.CUDAPlace(0)
414+
else:
415+
place = core.CPUPlace()
416+
with base.dygraph.guard(place):
417+
input_data = np.array(
418+
[
419+
[
420+
[
421+
[1, 2, 3, 4, 5],
422+
[6, 7, 8, 9, 10],
423+
[11, 12, 13, 14, 15],
424+
[16, 17, 18, 19, 20],
425+
]
426+
]
427+
]
428+
).astype("float32")
429+
input_x = paddle.to_tensor(input_data)
430+
output, indices = F.max_pool2d(
431+
input_x, kernel_size=2, stride=2, return_mask=True
432+
)
433+
out_pp = F.max_unpool2d(
434+
output.astype("int64"),
435+
indices,
436+
kernel_size=2,
437+
stride=None,
438+
output_size=input_x.shape,
439+
)
440+
output_np = output.numpy()
441+
indices_np = indices.numpy()
442+
expect_res = unpool2dmax_forward_naive(
443+
output_np, indices_np, [2, 2], [2, 2], [0, 0], [4, 5]
444+
).astype("float64")
445+
np.testing.assert_allclose(out_pp.numpy(), expect_res, rtol=1e-05)
446+
447+
403448
class TestUnpoolOpAPI_st(unittest.TestCase):
404449
@test_with_pir_api
405450
def test_case(self):

0 commit comments

Comments
 (0)