Skip to content

Commit

Permalink
remove fp32 tmp tensor and cast op for initializer.Normal and initial…
Browse files Browse the repository at this point in the history
…izer.Constant (PaddlePaddle#38818)
  • Loading branch information
GuoxiaWang authored Jan 10, 2022
1 parent 04f73d8 commit 2238a53
Show file tree
Hide file tree
Showing 4 changed files with 42 additions and 100 deletions.
17 changes: 12 additions & 5 deletions paddle/fluid/operators/gaussian_random_op.cu
Original file line number Diff line number Diff line change
Expand Up @@ -18,6 +18,7 @@ limitations under the License. */
#include "paddle/fluid/framework/generator.h"
#include "paddle/fluid/framework/op_registry.h"
#include "paddle/fluid/framework/operator.h"
#include "paddle/fluid/operators/amp/fp16_type_traits.h"
#include "paddle/fluid/operators/fill_constant_op.h"

namespace paddle {
Expand All @@ -38,10 +39,12 @@ struct GaussianGenerator {
__host__ __device__ T operator()(const unsigned int n) const {
thrust::minstd_rand rng;
rng.seed(seed_);
thrust::normal_distribution<T> dist(mean_, std_);
using MT = typename details::MPTypeTrait<T>::Type;
thrust::normal_distribution<MT> dist(mean_, std_);
unsigned int new_n = n + offset_;
rng.discard(new_n);
return dist(rng);
MT out = dist(rng);
return static_cast<T>(out);
}
};

Expand Down Expand Up @@ -124,10 +127,14 @@ class GPUGaussianRandomBatchSizeLikeKernel : public framework::OpKernel<T> {
} // namespace operators
} // namespace paddle

REGISTER_OP_CUDA_KERNEL(gaussian_random,
paddle::operators::GPUGaussianRandomKernel<float>,
paddle::operators::GPUGaussianRandomKernel<double>);
REGISTER_OP_CUDA_KERNEL(
gaussian_random,
paddle::operators::GPUGaussianRandomKernel<paddle::platform::float16>,
paddle::operators::GPUGaussianRandomKernel<float>,
paddle::operators::GPUGaussianRandomKernel<double>);
REGISTER_OP_CUDA_KERNEL(
gaussian_random_batch_size_like,
paddle::operators::GPUGaussianRandomBatchSizeLikeKernel<
paddle::platform::float16>,
paddle::operators::GPUGaussianRandomBatchSizeLikeKernel<float>,
paddle::operators::GPUGaussianRandomBatchSizeLikeKernel<double>);
99 changes: 21 additions & 78 deletions python/paddle/fluid/initializer.py
Original file line number Diff line number Diff line change
Expand Up @@ -137,54 +137,27 @@ def __call__(self, var, block=None):
isinstance(var, framework.EagerParamBase))
assert isinstance(block, framework.Block)

# to be compatible of fp16 initializers
if var.dtype == VarDesc.VarType.FP16:
out_dtype = VarDesc.VarType.FP32
out_var = block.create_var(
name=unique_name.generate(".".join(
['constant_init', var.name, 'tmp'])),
shape=var.shape,
dtype=out_dtype,
type=VarDesc.VarType.LOD_TENSOR,
persistable=False)
else:
out_dtype = var.dtype
out_var = var

if framework.in_dygraph_mode():
out_var = _C_ops.fill_constant(
out_var, 'value',
var = _C_ops.fill_constant(
var, 'value',
float(self._value), 'force_cpu', self._force_cpu, 'dtype',
int(out_dtype), 'str_value',
int(var.dtype), 'str_value',
str(float(self._value)), 'shape', var.shape)
if var.dtype == VarDesc.VarType.FP16:
var_tmp = _C_ops.cast(out_var, 'in_dtype', out_var.dtype,
'out_dtype', var.dtype)
var.copy_(var_tmp, False)
else:
var.copy_(out_var, False)
return None
else:
# fill constant should set the "str_value" to preserve precision
op = block.append_op(
type="fill_constant",
outputs={"Out": out_var},
outputs={"Out": var},
attrs={
"shape": var.shape,
"dtype": int(out_dtype),
"dtype": int(var.dtype),
"value": float(self._value),
'str_value': str(float(self._value)),
'force_cpu': self._force_cpu
},
stop_gradient=True)

if var.dtype == VarDesc.VarType.FP16:
block.append_op(
type="cast",
inputs={"X": out_var},
outputs={"Out": var},
attrs={"in_dtype": out_var.dtype,
"out_dtype": var.dtype})
var.op = op
return op

Expand Down Expand Up @@ -361,54 +334,24 @@ def __call__(self, var, block=None):
if self._seed == 0:
self._seed = block.program.random_seed

# to be compatible of fp16 initalizers
if var.dtype in [VarDesc.VarType.FP16, VarDesc.VarType.BF16]:
out_dtype = VarDesc.VarType.FP32
out_var = block.create_var(
name=unique_name.generate(".".join(
['gaussian_random', var.name, 'tmp'])),
shape=var.shape,
dtype=out_dtype,
type=VarDesc.VarType.LOD_TENSOR,
persistable=False)
else:
out_dtype = var.dtype
out_var = var

if framework.in_dygraph_mode():
out_var = _C_ops.gaussian_random(
'shape', var.shape, 'dtype', out_dtype, 'mean', self._mean,
'std', self._std_dev, 'seed', self._seed, 'use_mkldnn', False)
if var.dtype in [VarDesc.VarType.FP16, VarDesc.VarType.BF16]:
var_tmp = _C_ops.cast(out_var, 'in_dtype', out_var.dtype,
'out_dtype', var.dtype)
var.copy_(var_tmp, False)
else:
var.copy_(out_var, False)
return None
else:
op = block.append_op(
type="gaussian_random",
outputs={"Out": out_var},
attrs={
"shape": var.shape,
"dtype": out_dtype,
"mean": self._mean,
"std": self._std_dev,
"seed": self._seed,
"use_mkldnn": False
},
stop_gradient=True)

if var.dtype in [VarDesc.VarType.FP16, VarDesc.VarType.BF16]:
block.append_op(
type="cast",
inputs={"X": out_var},
outputs={"Out": var},
attrs={"in_dtype": out_var.dtype,
"out_dtype": var.dtype})
op = block.append_op(
type="gaussian_random",
outputs={"Out": var},
attrs={
"shape": var.shape,
"dtype": var.dtype,
"mean": self._mean,
"std": self._std_dev,
"seed": self._seed,
"use_mkldnn": False
},
stop_gradient=True)

if not framework.in_dygraph_mode():
var.op = op
return op
else:
return None


class TruncatedNormalInitializer(Initializer):
Expand Down
18 changes: 7 additions & 11 deletions python/paddle/fluid/tests/unittests/test_initializer.py
Original file line number Diff line number Diff line change
Expand Up @@ -65,7 +65,7 @@ def test_constant_initializer_default_value(self, dtype="float32"):
lod_level=0,
name="param",
initializer=initializer.ConstantInitializer())
num_ops = 2 if dtype == "float16" else 1
num_ops = 1
self.assertEqual(len(block.ops), num_ops)
init_op = block.ops[0]
self.assertEqual(init_op.type, 'fill_constant')
Expand All @@ -84,7 +84,7 @@ def test_constant_initializer(self, dtype="float32"):
lod_level=0,
name="param",
initializer=initializer.ConstantInitializer(2.3))
num_ops = 2 if dtype == "float16" else 1
num_ops = 1
self.assertEqual(len(block.ops), num_ops)
init_op = block.ops[0]
self.assertEqual(init_op.type, 'fill_constant')
Expand All @@ -94,10 +94,8 @@ def test_constant_initializer(self, dtype="float32"):
def test_constant_initializer_fp16(self):
"""Test constant initializer with float16
"""
block = self.test_constant_initializer_default_value("float16")
self.assertTrue(check_cast_op(block.ops[1]))
block = self.test_constant_initializer("float16")
self.assertTrue(check_cast_op(block.ops[1]))
self.test_constant_initializer_default_value("float16")
self.test_constant_initializer("float16")

def test_constant_initializer_bf16(self):
"""Test constant initializer with bfloat16
Expand Down Expand Up @@ -246,7 +244,7 @@ def test_normal_initializer(self, dtype="float32"):
lod_level=0,
name="param",
initializer=initializer.NormalInitializer(2.3, 1.9, 123))
num_ops = 2 if dtype in ["float16", "uint16"] else 1
num_ops = 1
self.assertEqual(len(block.ops), num_ops)
init_op = block.ops[0]
self.assertEqual(init_op.type, 'gaussian_random')
Expand All @@ -258,14 +256,12 @@ def test_normal_initializer(self, dtype="float32"):
def test_normal_initializer_fp16(self):
"""Test normal initializer with float16
"""
block = self.test_normal_initializer("float16")
self.assertTrue(check_cast_op(block.ops[1]))
self.test_normal_initializer("float16")

def test_normal_initializer_bf16(self):
"""Test normal initializer with bfloat16
"""
block = self.test_normal_initializer("uint16")
self.assertTrue(check_cast_op(block.ops[1]))
self.test_normal_initializer("uint16")


class TestXavierInitializer(unittest.TestCase):
Expand Down
8 changes: 2 additions & 6 deletions python/paddle/fluid/tests/unittests/test_initializer_nn.py
Original file line number Diff line number Diff line change
Expand Up @@ -54,7 +54,7 @@ def static_test_constant_initializer_common(self,
lod_level=0,
name="param",
initializer=init_inst)
num_ops = 2 if dtype in ["float16"] else 1
num_ops = 1
self.assertEqual(len(block.ops), num_ops)
init_op = block.ops[0]
self.assertEqual(init_op.type, 'fill_constant')
Expand Down Expand Up @@ -103,9 +103,7 @@ def test_constant_initializer_fp16(self):
"""Test constant initializer with float16
"""
block = self.test_constant_initializer_default_value_static("float16")
self.assertTrue(check_cast_op(block.ops[1]))
block = self.test_constant_initializer_static("float16")
self.assertTrue(check_cast_op(block.ops[1]))
self.test_constant_initializer_default_value_dygraph("float16")
self.test_constant_initializer_dygraph("float16")

Expand Down Expand Up @@ -402,7 +400,7 @@ def test_normal_initializer(self, dtype="float32"):
lod_level=0,
name="param",
initializer=initializer.Normal(2.3, 1.9))
num_ops = 2 if dtype in ["float16", "uint16"] else 1
num_ops = 1
self.assertEqual(len(block.ops), num_ops)
init_op = block.ops[0]
self.assertEqual(init_op.type, 'gaussian_random')
Expand All @@ -417,13 +415,11 @@ def test_normal_initializer_fp16(self):
"""Test normal initializer with float16
"""
block = self.test_normal_initializer("float16")
self.assertTrue(check_cast_op(block.ops[1]))

def test_normal_initializer_bf16(self):
"""Test normal initializer with bfloat16
"""
block = self.test_normal_initializer("uint16") #bfloat16
self.assertTrue(check_cast_op(block.ops[1]))

def test_normal_initializer_dygraph(self):
"""Test normal initializer in dygraph model.
Expand Down

0 comments on commit 2238a53

Please sign in to comment.