Skip to content

Commit 0be3494

Browse files
author
chengduo
authored
Merge pull request #3594 from chengduoZH/Adapting_to_the_BatchNorm_structure_to_support_3D_data
Adapting to the BatchNorm structure to support 3D data
2 parents 2fefaa1 + 5500153 commit 0be3494

File tree

15 files changed

+354
-16
lines changed

15 files changed

+354
-16
lines changed

paddle/gserver/layers/BatchNormBaseLayer.cpp

Lines changed: 5 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -62,14 +62,18 @@ void BatchNormBaseLayer::calFeatureMapSize() {
6262
const ImageConfig& conf = config_.inputs(0).image_conf();
6363
imageH_ = inputLayers_[0]->getOutput().getFrameHeight();
6464
imageW_ = inputLayers_[0]->getOutput().getFrameWidth();
65+
imageD_ = inputLayers_[0]->getOutput().getFrameDepth();
66+
67+
if (0 == imageD_) imageD_ = conf.img_size_z();
6568
if (imageH_ == 0 && imageW_ == 0) {
6669
imageH_ = conf.has_img_size_y() ? conf.img_size_y() : conf.img_size();
6770
imageW_ = conf.img_size();
6871
} else {
6972
getOutput().setFrameHeight(imageH_);
7073
getOutput().setFrameWidth(imageW_);
74+
getOutput().setFrameDepth(imageD_);
7175
}
72-
imgPixels_ = imageH_ * imageW_;
76+
imgPixels_ = imageH_ * imageW_ * imageD_;
7377
}
7478

7579
} // namespace paddle

paddle/gserver/layers/BatchNormBaseLayer.h

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -80,6 +80,7 @@ class BatchNormBaseLayer : public Layer {
8080

8181
/// Height or width of input image feature.
8282
/// Both of them are 1 if the input is fully-connected layer.
83+
int imageD_;
8384
int imageH_;
8485
int imageW_;
8586
/// Height * Width.

paddle/gserver/layers/CudnnBatchNormLayer.cpp

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -37,7 +37,7 @@ bool CudnnBatchNormLayer::init(const LayerMap& layerMap,
3737
}
3838

3939
void CudnnBatchNormLayer::reshape(int batchSize) {
40-
hl_tensor_reshape(ioDesc_, batchSize, channels_, imageH_, imageW_);
40+
hl_tensor_reshape(ioDesc_, batchSize, channels_, imageH_ * imageD_, imageW_);
4141
}
4242

4343
void CudnnBatchNormLayer::forward(PassType passType) {
@@ -104,7 +104,7 @@ void CudnnBatchNormLayer::forward(PassType passType) {
104104
EPS,
105105
batchSize,
106106
channels_,
107-
imageH_,
107+
imageH_ * imageD_,
108108
imageW_);
109109
}
110110
}

paddle/gserver/tests/test_LayerGrad.cpp

Lines changed: 49 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1703,6 +1703,55 @@ TEST(Layer, BatchNormalizationLayer) {
17031703
#endif
17041704
}
17051705

1706+
void testBatchNorm3DLayer(const string& type, bool trans, bool useGpu) {
1707+
TestConfig config;
1708+
const int CHANNELS = 10;
1709+
const int IMG_SIZE = 16;
1710+
const int IMG_SIZE_Y = 8;
1711+
const int IMG_SIZE_Z = 8;
1712+
size_t size = CHANNELS * IMG_SIZE * IMG_SIZE_Y * IMG_SIZE_Z;
1713+
config.layerConfig.set_type(type);
1714+
config.layerConfig.set_size(size);
1715+
config.layerConfig.set_active_type("sigmoid");
1716+
config.biasSize = CHANNELS;
1717+
config.inputDefs.push_back({INPUT_DATA,
1718+
"layer_0",
1719+
/* dim= */ size,
1720+
/* paraSize= */ CHANNELS});
1721+
1722+
config.inputDefs.push_back({INPUT_DATA, "layer_1_running_mean", 1, CHANNELS});
1723+
config.inputDefs.back().isStatic = true;
1724+
config.inputDefs.push_back({INPUT_DATA, "layer_2_running_var", 1, CHANNELS});
1725+
config.inputDefs.back().isStatic = true;
1726+
1727+
LayerInputConfig* input = config.layerConfig.add_inputs();
1728+
config.layerConfig.add_inputs();
1729+
config.layerConfig.add_inputs();
1730+
1731+
ImageConfig* img_conf = input->mutable_image_conf();
1732+
img_conf->set_channels(CHANNELS);
1733+
img_conf->set_img_size(IMG_SIZE);
1734+
img_conf->set_img_size_y(IMG_SIZE_Y);
1735+
img_conf->set_img_size_z(IMG_SIZE_Z);
1736+
1737+
testLayerGrad(config,
1738+
"batch_norm",
1739+
64,
1740+
/* trans= */ trans,
1741+
useGpu,
1742+
/* useWeight */ true);
1743+
}
1744+
1745+
TEST(Layer, testBatchNorm3DLayer) {
1746+
testBatchNorm3DLayer("batch_norm", false, false);
1747+
#ifndef PADDLE_ONLY_CPU
1748+
testBatchNorm3DLayer("batch_norm", false, true);
1749+
if (hl_get_cudnn_lib_version() >= int(4000)) {
1750+
testBatchNorm3DLayer("cudnn_batch_norm", false, true);
1751+
}
1752+
#endif
1753+
}
1754+
17061755
void testConvOperator(bool isDeconv) {
17071756
TestConfig config;
17081757
const int NUM_FILTERS = 16;

proto/ModelConfig.proto

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -520,6 +520,7 @@ message LayerConfig {
520520
// for HuberRegressionLoss
521521
optional double delta = 57 [ default = 1.0 ];
522522

523+
// for 3D data
523524
optional uint64 depth = 58 [ default = 1 ];
524525

525526
// for switch order layer

python/paddle/trainer/config_parser.py

Lines changed: 78 additions & 11 deletions
Original file line numberDiff line numberDiff line change
@@ -1332,6 +1332,12 @@ def parse_image(image, input_layer_name, image_conf):
13321332
get_img_size(input_layer_name, image_conf.channels)
13331333

13341334

1335+
def parse_image3d(image, input_layer_name, image_conf):
1336+
image_conf.channels = image.channels
1337+
image_conf.img_size, image_conf.img_size_y, image_conf.img_size_z = \
1338+
get_img3d_size(input_layer_name, image_conf.channels)
1339+
1340+
13351341
def parse_norm(norm, input_layer_name, norm_conf):
13361342
norm_conf.norm_type = norm.norm_type
13371343
config_assert(
@@ -2365,6 +2371,7 @@ def __init__(self,
23652371
name,
23662372
inputs,
23672373
bias=True,
2374+
img3D=False,
23682375
use_global_stats=True,
23692376
moving_average_fraction=0.9,
23702377
batch_norm_type=None,
@@ -2410,15 +2417,33 @@ def __init__(self,
24102417

24112418
input_layer = self.get_input_layer(0)
24122419
image_conf = self.config.inputs[0].image_conf
2413-
parse_image(self.inputs[0].image, input_layer.name, image_conf)
2414-
2415-
# Only pass the width and height of input to batch_norm layer
2416-
# when either of it is non-zero.
2417-
if input_layer.width != 0 or input_layer.height != 0:
2418-
self.set_cnn_layer(name, image_conf.img_size_y, image_conf.img_size,
2419-
image_conf.channels, False)
2420+
if img3D:
2421+
parse_image3d(self.inputs[0].image, input_layer.name, image_conf)
2422+
# Only pass the width and height of input to batch_norm layer
2423+
# when either of it is non-zero.
2424+
if input_layer.width != 0 or input_layer.height != 0:
2425+
self.set_cnn_layer(
2426+
input_layer_name=name,
2427+
depth=image_conf.img_size_z,
2428+
height=image_conf.img_size_y,
2429+
width=image_conf.img_size,
2430+
channels=image_conf.channels,
2431+
is_print=True)
2432+
else:
2433+
self.set_layer_size(input_layer.size)
24202434
else:
2421-
self.set_layer_size(input_layer.size)
2435+
parse_image(self.inputs[0].image, input_layer.name, image_conf)
2436+
# Only pass the width and height of input to batch_norm layer
2437+
# when either of it is non-zero.
2438+
if input_layer.width != 0 or input_layer.height != 0:
2439+
self.set_cnn_layer(
2440+
input_layer_name=name,
2441+
height=image_conf.img_size_y,
2442+
width=image_conf.img_size,
2443+
channels=image_conf.channels,
2444+
is_print=True)
2445+
else:
2446+
self.set_layer_size(input_layer.size)
24222447

24232448
psize = self.calc_parameter_size(image_conf)
24242449
dims = [1, psize]
@@ -2433,6 +2458,28 @@ def __init__(self,
24332458

24342459
self.create_bias_parameter(bias, psize)
24352460

2461+
def set_cnn_layer(self,
2462+
input_layer_name,
2463+
depth=None,
2464+
height=None,
2465+
width=None,
2466+
channels=None,
2467+
is_print=True):
2468+
depthIsNone = False
2469+
if depth is None:
2470+
depth = 1
2471+
depthIsNone = True
2472+
size = depth * height * width * channels
2473+
self.set_layer_size(size)
2474+
self.set_layer_height_width(height, width)
2475+
self.set_layer_depth(depth)
2476+
if is_print and depthIsNone:
2477+
print("output for %s: c = %d, h = %d, w = %d, size = %d" %
2478+
(input_layer_name, channels, height, width, size))
2479+
elif is_print:
2480+
print("output for %s: c = %d, d = %d, h = %d, w = %d, size = %d" %
2481+
(input_layer_name, channels, depth, height, width, size))
2482+
24362483
def calc_parameter_size(self, image_conf):
24372484
return image_conf.channels
24382485

@@ -2694,9 +2741,20 @@ def __init__(self, name, inputs, bias=True, **xargs):
26942741
super(AddToLayer, self).__init__(
26952742
name, 'addto', 0, inputs=inputs, **xargs)
26962743
config_assert(len(inputs) > 0, 'inputs cannot be empty for AddToLayer')
2697-
for input_index in xrange(len(self.inputs)):
2698-
input_layer = self.get_input_layer(input_index)
2699-
self.set_layer_size(input_layer.size)
2744+
2745+
if len(self.inputs) > 1:
2746+
for input_index in xrange(len(self.inputs)):
2747+
assert self.get_input_layer(0).height == self.get_input_layer(
2748+
input_index).height
2749+
assert self.get_input_layer(0).width == self.get_input_layer(
2750+
input_index).width
2751+
assert self.get_input_layer(0).depth == self.get_input_layer(
2752+
input_index).depth
2753+
2754+
self.set_layer_size(self.get_input_layer(0).size)
2755+
self.set_layer_height_width(self.get_input_layer(0).height, \
2756+
self.get_input_layer(0).width)
2757+
self.set_layer_depth(self.get_input_layer(0).depth)
27002758
self.create_bias_parameter(bias, self.config.size)
27012759

27022760

@@ -3376,11 +3434,20 @@ def __init__(self, name, inputs, bias=False, **xargs):
33763434
name, 'concat', 0, inputs=inputs, **xargs)
33773435
size = 0
33783436
for input_index in xrange(len(self.inputs)):
3437+
assert self.get_input_layer(0).height == self.get_input_layer(
3438+
input_index).height
3439+
assert self.get_input_layer(0).width == self.get_input_layer(
3440+
input_index).width
3441+
assert self.get_input_layer(0).depth == self.get_input_layer(
3442+
input_index).depth
33793443
input_layer = self.get_input_layer(input_index)
33803444
input = self.inputs[input_index]
33813445
if self.config.size == 0:
33823446
size += input_layer.size
33833447

3448+
self.set_layer_height_width(self.get_input_layer(0).height, \
3449+
self.get_input_layer(0).width)
3450+
self.set_layer_depth(self.get_input_layer(0).depth)
33843451
self.set_layer_size(size)
33853452

33863453

python/paddle/trainer_config_helpers/layers.py

Lines changed: 7 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -354,6 +354,10 @@ def width(self):
354354
def height(self):
355355
return cp.g_layer_map[self.full_name].height
356356

357+
@property
358+
def depth(self):
359+
return cp.g_layer_map[self.full_name].depth
360+
357361
def set_input(self, input):
358362
"""
359363
Set the input for a memory layer. Can only be used for memory layer
@@ -943,7 +947,7 @@ def data_layer(name, size, depth=None, height=None, width=None,
943947
if height is not None and width is not None:
944948
num_filters = size / (width * height * depth)
945949
assert num_filters * width * height * depth == size, \
946-
"size=%s width=%s height=%s depth=%s" % (size, width, height, depth)
950+
"size=%s width=%s height=%s depth=%s" % (size, width, height, depth)
947951

948952
return LayerOutput(name, LayerType.DATA, size=size, num_filters=num_filters)
949953

@@ -2953,6 +2957,7 @@ def img_cmrnorm_layer(input,
29532957
def batch_norm_layer(input,
29542958
act=None,
29552959
name=None,
2960+
img3D=False,
29562961
num_channels=None,
29572962
bias_attr=None,
29582963
param_attr=None,
@@ -3042,6 +3047,7 @@ def batch_norm_layer(input,
30423047
(batch_norm_type == "cudnn_batch_norm")
30433048
l = Layer(
30443049
name=name,
3050+
img3D=img3D,
30453051
inputs=Input(
30463052
input.name, image=Image(channels=num_channels), **param_attr.attr),
30473053
active_type=act.name,

python/paddle/trainer_config_helpers/tests/configs/file_list.sh

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -10,6 +10,6 @@ test_prelu_layer test_row_conv test_detection_output_layer test_multibox_loss_la
1010
test_recursive_topology test_gated_unit_layer test_clip_layer test_row_l2_norm_layer
1111
test_kmax_seq_socre_layer test_sub_nested_seq_select_layer test_scale_shift_layer
1212
test_seq_slice_layer test_cross_entropy_over_beam test_pooling3D_layer
13-
test_conv3d_layer test_deconv3d_layer)
13+
test_conv3d_layer test_deconv3d_layer test_BatchNorm3D)
1414

1515
export whole_configs=(test_split_datasource)

python/paddle/trainer_config_helpers/tests/configs/protostr/img_layers.protostr

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -62,6 +62,7 @@ layers {
6262
moving_average_fraction: 0.9
6363
height: 227
6464
width: 227
65+
depth: 1
6566
}
6667
layers {
6768
name: "__crmnorm_0__"

python/paddle/trainer_config_helpers/tests/configs/protostr/img_trans_layers.protostr

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -62,6 +62,7 @@ layers {
6262
moving_average_fraction: 0.9
6363
height: 256
6464
width: 256
65+
depth: 1
6566
}
6667
layers {
6768
name: "__crmnorm_0__"

0 commit comments

Comments
 (0)