-
Notifications
You must be signed in to change notification settings - Fork 18.7k
/
test_lrn_layer.cpp
282 lines (260 loc) · 10.3 KB
/
test_lrn_layer.cpp
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
// Copyright 2014 BVLC and contributors.
#include <algorithm>
#include <cstring>
#include <vector>
#include "cuda_runtime.h"
#include "gtest/gtest.h"
#include "caffe/blob.hpp"
#include "caffe/common.hpp"
#include "caffe/filler.hpp"
#include "caffe/vision_layers.hpp"
#include "caffe/test/test_gradient_check_util.hpp"
#include "caffe/test/test_caffe_main.hpp"
using std::min;
using std::max;
namespace caffe {
extern cudaDeviceProp CAFFE_TEST_CUDA_PROP;
template <typename Dtype>
class LRNLayerTest : public ::testing::Test {
protected:
LRNLayerTest()
: blob_bottom_(new Blob<Dtype>()),
blob_top_(new Blob<Dtype>()),
epsilon_(Dtype(1e-5)) {}
virtual void SetUp() {
Caffe::set_random_seed(1701);
blob_bottom_->Reshape(2, 7, 3, 3);
// fill the values
FillerParameter filler_param;
GaussianFiller<Dtype> filler(filler_param);
filler.Fill(this->blob_bottom_);
blob_bottom_vec_.push_back(blob_bottom_);
blob_top_vec_.push_back(blob_top_);
}
virtual ~LRNLayerTest() { delete blob_bottom_; delete blob_top_; }
void ReferenceLRNForward(const Blob<Dtype>& blob_bottom,
const LayerParameter& layer_param, Blob<Dtype>* blob_top);
Dtype epsilon_;
Blob<Dtype>* const blob_bottom_;
Blob<Dtype>* const blob_top_;
vector<Blob<Dtype>*> blob_bottom_vec_;
vector<Blob<Dtype>*> blob_top_vec_;
};
template <typename Dtype>
void LRNLayerTest<Dtype>::ReferenceLRNForward(
const Blob<Dtype>& blob_bottom, const LayerParameter& layer_param,
Blob<Dtype>* blob_top) {
blob_top->Reshape(blob_bottom.num(), blob_bottom.channels(),
blob_bottom.height(), blob_bottom.width());
const Dtype* bottom_data = blob_bottom.cpu_data();
Dtype* top_data = blob_top->mutable_cpu_data();
LRNParameter lrn_param = layer_param.lrn_param();
Dtype alpha = lrn_param.alpha();
Dtype beta = lrn_param.beta();
int size = lrn_param.local_size();
switch (lrn_param.norm_region()) {
case LRNParameter_NormRegion_ACROSS_CHANNELS:
for (int n = 0; n < blob_bottom.num(); ++n) {
for (int c = 0; c < blob_bottom.channels(); ++c) {
for (int h = 0; h < blob_bottom.height(); ++h) {
for (int w = 0; w < blob_bottom.width(); ++w) {
int c_start = c - (size - 1) / 2;
int c_end = min(c_start + size, blob_bottom.channels());
c_start = max(c_start, 0);
Dtype scale = 1.;
for (int i = c_start; i < c_end; ++i) {
Dtype value = blob_bottom.data_at(n, i, h, w);
scale += value * value * alpha / size;
}
*(top_data + blob_top->offset(n, c, h, w)) =
blob_bottom.data_at(n, c, h, w) / pow(scale, beta);
}
}
}
}
break;
case LRNParameter_NormRegion_WITHIN_CHANNEL:
for (int n = 0; n < blob_bottom.num(); ++n) {
for (int c = 0; c < blob_bottom.channels(); ++c) {
for (int h = 0; h < blob_bottom.height(); ++h) {
int h_start = h - (size - 1) / 2;
int h_end = min(h_start + size, blob_bottom.height());
h_start = max(h_start, 0);
for (int w = 0; w < blob_bottom.width(); ++w) {
Dtype scale = 1.;
int w_start = w - (size - 1) / 2;
int w_end = min(w_start + size, blob_bottom.width());
w_start = max(w_start, 0);
for (int nh = h_start; nh < h_end; ++nh) {
for (int nw = w_start; nw < w_end; ++nw) {
Dtype value = blob_bottom.data_at(n, c, nh, nw);
scale += value * value * alpha / (size * size);
}
}
*(top_data + blob_top->offset(n, c, h, w)) =
blob_bottom.data_at(n, c, h, w) / pow(scale, beta);
}
}
}
}
break;
default:
LOG(FATAL) << "Unknown normalization region.";
}
}
typedef ::testing::Types<float, double> Dtypes;
TYPED_TEST_CASE(LRNLayerTest, Dtypes);
TYPED_TEST(LRNLayerTest, TestSetupAcrossChannels) {
LayerParameter layer_param;
LRNLayer<TypeParam> layer(layer_param);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
EXPECT_EQ(this->blob_top_->num(), 2);
EXPECT_EQ(this->blob_top_->channels(), 7);
EXPECT_EQ(this->blob_top_->height(), 3);
EXPECT_EQ(this->blob_top_->width(), 3);
}
TYPED_TEST(LRNLayerTest, TestCPUForwardAcrossChannels) {
LayerParameter layer_param;
LRNLayer<TypeParam> layer(layer_param);
Caffe::set_mode(Caffe::CPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
Blob<TypeParam> top_reference;
this->ReferenceLRNForward(*(this->blob_bottom_), layer_param,
&top_reference);
for (int i = 0; i < this->blob_bottom_->count(); ++i) {
EXPECT_NEAR(this->blob_top_->cpu_data()[i], top_reference.cpu_data()[i],
this->epsilon_);
}
}
TYPED_TEST(LRNLayerTest, TestGPUForwardAcrossChannels) {
LayerParameter layer_param;
LRNLayer<TypeParam> layer(layer_param);
Caffe::set_mode(Caffe::GPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
Blob<TypeParam> top_reference;
this->ReferenceLRNForward(*(this->blob_bottom_), layer_param,
&top_reference);
for (int i = 0; i < this->blob_bottom_->count(); ++i) {
EXPECT_NEAR(this->blob_top_->cpu_data()[i], top_reference.cpu_data()[i],
this->epsilon_);
}
}
TYPED_TEST(LRNLayerTest, TestCPUGradientAcrossChannels) {
LayerParameter layer_param;
LRNLayer<TypeParam> layer(layer_param);
GradientChecker<TypeParam> checker(1e-2, 1e-2);
Caffe::set_mode(Caffe::CPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
for (int i = 0; i < this->blob_top_->count(); ++i) {
this->blob_top_->mutable_cpu_diff()[i] = 1.;
}
layer.Backward(this->blob_top_vec_, true, &(this->blob_bottom_vec_));
// for (int i = 0; i < this->blob_bottom_->count(); ++i) {
// std::cout << "CPU diff " << this->blob_bottom_->cpu_diff()[i]
// << std::endl;
// }
checker.CheckGradientExhaustive(&layer, &(this->blob_bottom_vec_),
&(this->blob_top_vec_));
}
TYPED_TEST(LRNLayerTest, TestGPUGradientAcrossChannels) {
LayerParameter layer_param;
LRNLayer<TypeParam> layer(layer_param);
GradientChecker<TypeParam> checker(1e-2, 1e-2);
Caffe::set_mode(Caffe::GPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
for (int i = 0; i < this->blob_top_->count(); ++i) {
this->blob_top_->mutable_cpu_diff()[i] = 1.;
}
layer.Backward(this->blob_top_vec_, true, &(this->blob_bottom_vec_));
// for (int i = 0; i < this->blob_bottom_->count(); ++i) {
// std::cout << "GPU diff " << this->blob_bottom_->cpu_diff()[i]
// << std::endl;
// }
checker.CheckGradientExhaustive(&layer, &(this->blob_bottom_vec_),
&(this->blob_top_vec_));
}
TYPED_TEST(LRNLayerTest, TestSetupWithinChannel) {
LayerParameter layer_param;
layer_param.mutable_lrn_param()->set_norm_region(
LRNParameter_NormRegion_WITHIN_CHANNEL);
layer_param.mutable_lrn_param()->set_local_size(3);
LRNLayer<TypeParam> layer(layer_param);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
EXPECT_EQ(this->blob_top_->num(), 2);
EXPECT_EQ(this->blob_top_->channels(), 7);
EXPECT_EQ(this->blob_top_->height(), 3);
EXPECT_EQ(this->blob_top_->width(), 3);
}
TYPED_TEST(LRNLayerTest, TestCPUForwardWithinChannel) {
LayerParameter layer_param;
layer_param.mutable_lrn_param()->set_norm_region(
LRNParameter_NormRegion_WITHIN_CHANNEL);
layer_param.mutable_lrn_param()->set_local_size(3);
LRNLayer<TypeParam> layer(layer_param);
Caffe::set_mode(Caffe::CPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
Blob<TypeParam> top_reference;
this->ReferenceLRNForward(*(this->blob_bottom_), layer_param,
&top_reference);
for (int i = 0; i < this->blob_bottom_->count(); ++i) {
EXPECT_NEAR(this->blob_top_->cpu_data()[i], top_reference.cpu_data()[i],
this->epsilon_);
}
}
TYPED_TEST(LRNLayerTest, TestGPUForwardWithinChannel) {
LayerParameter layer_param;
layer_param.mutable_lrn_param()->set_norm_region(
LRNParameter_NormRegion_WITHIN_CHANNEL);
layer_param.mutable_lrn_param()->set_local_size(3);
LRNLayer<TypeParam> layer(layer_param);
Caffe::set_mode(Caffe::GPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
Blob<TypeParam> top_reference;
this->ReferenceLRNForward(*(this->blob_bottom_), layer_param,
&top_reference);
for (int i = 0; i < this->blob_bottom_->count(); ++i) {
EXPECT_NEAR(this->blob_top_->cpu_data()[i], top_reference.cpu_data()[i],
this->epsilon_);
}
}
TYPED_TEST(LRNLayerTest, TestCPUGradientWithinChannel) {
LayerParameter layer_param;
layer_param.mutable_lrn_param()->set_norm_region(
LRNParameter_NormRegion_WITHIN_CHANNEL);
layer_param.mutable_lrn_param()->set_local_size(3);
LRNLayer<TypeParam> layer(layer_param);
GradientChecker<TypeParam> checker(1e-2, 1e-2);
Caffe::set_mode(Caffe::CPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
for (int i = 0; i < this->blob_top_->count(); ++i) {
this->blob_top_->mutable_cpu_diff()[i] = 1.;
}
layer.Backward(this->blob_top_vec_, true, &(this->blob_bottom_vec_));
checker.CheckGradientExhaustive(&layer, &(this->blob_bottom_vec_),
&(this->blob_top_vec_));
}
TYPED_TEST(LRNLayerTest, TestGPUGradientWithinChannel) {
LayerParameter layer_param;
layer_param.mutable_lrn_param()->set_norm_region(
LRNParameter_NormRegion_WITHIN_CHANNEL);
layer_param.mutable_lrn_param()->set_local_size(3);
LRNLayer<TypeParam> layer(layer_param);
GradientChecker<TypeParam> checker(1e-2, 1e-2);
Caffe::set_mode(Caffe::GPU);
layer.SetUp(this->blob_bottom_vec_, &(this->blob_top_vec_));
layer.Forward(this->blob_bottom_vec_, &(this->blob_top_vec_));
for (int i = 0; i < this->blob_top_->count(); ++i) {
this->blob_top_->mutable_cpu_diff()[i] = 1.;
}
layer.Backward(this->blob_top_vec_, true, &(this->blob_bottom_vec_));
checker.CheckGradientExhaustive(&layer, &(this->blob_bottom_vec_),
&(this->blob_top_vec_));
}
} // namespace caffe