Skip to content

Commit 4637151

Browse files
authored
add (N,C,*) input support for GroupNorm (#34773)
* add (N,C,*) input support for GroupNorm * --amend
1 parent 1aa2bde commit 4637151

File tree

5 files changed

+112
-27
lines changed

5 files changed

+112
-27
lines changed

paddle/fluid/operators/group_norm_op.cc

Lines changed: 7 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -37,6 +37,13 @@ class GroupNormOp : public framework::OperatorWithKernel {
3737
"GroupNorm");
3838

3939
auto x_dim = ctx->GetInputDim("X");
40+
PADDLE_ENFORCE_GE(
41+
x_dim.size(), 2,
42+
platform::errors::InvalidArgument(
43+
"The Input(X)'s dimension of Op(group_norm) must be "
44+
"greater than 1. But received: %u-D Tensor, which shape is [%s].",
45+
x_dim.size(), x_dim));
46+
4047
const std::string data_layout_str =
4148
ctx->Attrs().Get<std::string>("data_layout");
4249
const framework::DataLayout data_layout =

paddle/fluid/operators/group_norm_op.cu

Lines changed: 20 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -171,9 +171,16 @@ class GroupNormKernel<platform::CUDADeviceContext, T>
171171
const T* bias_data = nullptr;
172172
if (bias) bias_data = bias->data<T>();
173173

174-
int imsize = (data_layout == DataLayout::kNCHW ? x_dims[2] * x_dims[3]
175-
: x_dims[1] * x_dims[2]);
176-
174+
int imsize = 1;
175+
if (data_layout == DataLayout::kNCHW) {
176+
for (int i = 2; i < x_dims.size(); ++i) {
177+
imsize *= x_dims[i];
178+
}
179+
} else {
180+
for (int i = 1; i < x_dims.size() - 1; ++i) {
181+
imsize *= x_dims[i];
182+
}
183+
}
177184
#ifdef __HIPCC__
178185
int block_size = std::max(std::min(256, imsize), 64);
179186
#else
@@ -349,8 +356,16 @@ class GroupNormGradKernel<platform::CUDADeviceContext, T>
349356
const T* bias_data = nullptr;
350357
if (bias) bias_data = bias->data<T>();
351358

352-
int imsize = (data_layout == DataLayout::kNCHW ? x_dims[2] * x_dims[3]
353-
: x_dims[1] * x_dims[2]);
359+
int imsize = 1;
360+
if (data_layout == DataLayout::kNCHW) {
361+
for (int i = 2; i < x_dims.size(); ++i) {
362+
imsize *= x_dims[i];
363+
}
364+
} else {
365+
for (int i = 1; i < x_dims.size() - 1; ++i) {
366+
imsize *= x_dims[i];
367+
}
368+
}
354369

355370
#ifdef __HIPCC__
356371
int block_size = std::max(std::min(256, imsize), 64);

paddle/fluid/operators/group_norm_op.h

Lines changed: 20 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -68,9 +68,16 @@ class GroupNormKernel : public framework::OpKernel<T> {
6868
const T* bias_data = nullptr;
6969
if (bias) bias_data = bias->data<T>();
7070

71-
int imsize = (data_layout == DataLayout::kNCHW ? x_dims[2] * x_dims[3]
72-
: x_dims[1] * x_dims[2]);
73-
71+
int imsize = 1;
72+
if (data_layout == DataLayout::kNCHW) {
73+
for (int i = 2; i < x_dims.size(); ++i) {
74+
imsize *= x_dims[i];
75+
}
76+
} else {
77+
for (int i = 1; i < x_dims.size() - 1; ++i) {
78+
imsize *= x_dims[i];
79+
}
80+
}
7481
auto* iter_x_data = x_data;
7582
auto* iter_y_data = y_data;
7683
for (int bid = 0; bid < x_dims[0]; bid++) {
@@ -257,8 +264,16 @@ class GroupNormGradKernel : public framework::OpKernel<T> {
257264
const T* bias_data = nullptr;
258265
if (bias) bias_data = bias->data<T>();
259266

260-
int imsize = (data_layout == DataLayout::kNCHW ? x_dims[2] * x_dims[3]
261-
: x_dims[1] * x_dims[2]);
267+
int imsize = 1;
268+
if (data_layout == DataLayout::kNCHW) {
269+
for (int i = 2; i < x_dims.size(); ++i) {
270+
imsize *= x_dims[i];
271+
}
272+
} else {
273+
for (int i = 1; i < x_dims.size() - 1; ++i) {
274+
imsize *= x_dims[i];
275+
}
276+
}
262277
auto* iter_x_data = x_data;
263278
auto* iter_d_x_data = d_x_data;
264279
auto* iter_y_data = y_data;

python/paddle/fluid/tests/unittests/test_group_norm_op_v2.py

Lines changed: 63 additions & 15 deletions
Original file line numberDiff line numberDiff line change
@@ -25,13 +25,29 @@
2525
import paddle
2626

2727

28+
def group_norm_naive_for_general_dimension(x, scale, bias, epsilon, groups):
29+
# original version group norm only support 4-D tensor
30+
# this function generalizes to support differnt dimensions tensor (>= 2-D)
31+
input_shape = x.shape
32+
N, C = x.shape[0], x.shape[1]
33+
G = groups
34+
x = x.reshape((N * G, -1))
35+
mean = np.mean(x, axis=1, keepdims=True)
36+
var = np.var(x, axis=1, keepdims=True)
37+
output = (x - mean) / np.sqrt(var + epsilon)
38+
output = output.reshape(input_shape) * scale.reshape(
39+
(-1, 1, 1)) + bias.reshape((-1, 1, 1))
40+
return output
41+
42+
2843
class TestDygraphGroupNormv2(unittest.TestCase):
2944
def test_dygraph(self):
3045
places = [fluid.CPUPlace()]
3146
if core.is_compiled_with_cuda() and core.op_support_gpu("group_norm"):
3247
places.append(fluid.CUDAPlace(0))
48+
shapes = [[2, 2, 2, 2], [2, 2, 4], [4, 2], [4, 2, 6, 6, 2],
49+
[2, 2, 2, 2, 2, 2]]
3350
for p in places:
34-
shape = [2, 2, 2, 2]
3551

3652
def compute_v1(x):
3753
with fluid.dygraph.guard(p):
@@ -62,23 +78,26 @@ def attr_data_format():
6278

6379
self.assertRaises(ValueError, attr_data_format)
6480

65-
x = np.random.randn(*shape).astype("float32")
66-
y1 = compute_v1(x)
67-
y2 = compute_v2(x)
68-
result = np.allclose(y1, y2, atol=1e-5)
69-
if not result:
70-
print("y1:", y1, "\ty2:", y2)
71-
self.assertTrue(result)
72-
test_weight_bias_false()
73-
test_nn_exception()
81+
for shape in shapes:
82+
x = np.random.randn(*shape).astype("float32")
83+
y1 = compute_v1(x)
84+
y2 = compute_v2(x)
85+
result = np.allclose(y1, y2, atol=1e-5)
86+
if not result:
87+
print("y1:", y1, "\ty2:", y2)
88+
self.assertTrue(result)
89+
test_weight_bias_false()
90+
test_nn_exception()
7491

7592
def test_static(self):
93+
paddle.enable_static()
7694
places = [fluid.CPUPlace()]
7795
if core.is_compiled_with_cuda() and core.op_support_gpu("group_norm"):
7896
places.append(fluid.CUDAPlace(0))
97+
shapes = [[2, 6, 2, 2], [2, 6, 4], [4, 6], [4, 6, 6, 6, 2],
98+
[4, 6, 2, 2, 2, 2]]
7999
for p in places:
80100
exe = fluid.Executor(p)
81-
shape = [2, 6, 2, 2]
82101

83102
def compute_v1(x_np):
84103
with program_guard(Program(), Program()):
@@ -98,10 +117,39 @@ def compute_v2(x_np):
98117
r = exe.run(feed={'x': x_np}, fetch_list=[y])[0]
99118
return r
100119

101-
x = np.random.randn(*shape).astype("float32")
102-
y1 = compute_v1(x)
103-
y2 = compute_v2(x)
104-
self.assertTrue(np.allclose(y1, y2, atol=1e-5))
120+
for shape in shapes:
121+
x = np.random.randn(*shape).astype("float32")
122+
y1 = compute_v1(x)
123+
y2 = compute_v2(x)
124+
self.assertTrue(np.allclose(y1, y2, atol=1e-5))
125+
126+
127+
class TestGroupNormAPIV2_With_General_Dimensions(unittest.TestCase):
128+
def test_numerical_accuracy(self):
129+
paddle.disable_static()
130+
shapes = [(2, 6), (2, 6, 4), (2, 6, 4, 4), (2, 6, 6, 6, 2), (2, 6, 6, 6,
131+
2, 3)]
132+
places = [fluid.CPUPlace()]
133+
if core.is_compiled_with_cuda() and core.op_support_gpu("group_norm"):
134+
places.append(fluid.CUDAPlace(0))
135+
136+
for place in places:
137+
for shape in shapes:
138+
scale = np.array([1]).astype("float32")
139+
bias = np.array([0]).astype("float32")
140+
data = np.random.random(shape).astype("float32")
141+
expect_res1 = group_norm_naive_for_general_dimension(
142+
data, scale, bias, epsilon=1e-5, groups=6)
143+
expect_res2 = group_norm_naive_for_general_dimension(
144+
data, scale, bias, epsilon=1e-5, groups=2)
145+
146+
gn1 = paddle.nn.GroupNorm(num_channels=6, num_groups=6)
147+
gn2 = paddle.nn.GroupNorm(num_channels=6, num_groups=2)
148+
data_pd = paddle.to_tensor(data)
149+
result1 = gn1(data_pd).numpy()
150+
result2 = gn2(data_pd).numpy()
151+
self.assertTrue(np.allclose(result1, expect_res1, atol=1e-5))
152+
self.assertTrue(np.allclose(result2, expect_res2, atol=1e-5))
105153

106154

107155
if __name__ == '__main__':

python/paddle/nn/layer/norm.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -338,8 +338,8 @@ class GroupNorm(Layer):
338338
name(str, optional): Name for the GroupNorm, default is None. For more information, please refer to :ref:`api_guide_Name`..
339339
340340
Shape:
341-
- x: 4-D tensor with shape: (batch, num_features, height, weight).
342-
- output: 4-D tensor with same shape as input x.
341+
- x: Tensor with shape: (batch, num_features, *).
342+
- output: The same shape as input x.
343343
344344
Returns:
345345
None

0 commit comments

Comments
 (0)