Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Random op #3060

Merged
merged 37 commits into from
Aug 9, 2017
Merged
Show file tree
Hide file tree
Changes from 14 commits
Commits
Show all changes
37 commits
Select commit Hold shift + click to select a range
5ad9474
add random op
dzhwinter Jul 24, 2017
c110f56
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Jul 24, 2017
0d554f1
"add template fill function"
dzhwinter Jul 24, 2017
6f80b5f
"move to template function"
dzhwinter Jul 25, 2017
d263cce
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Jul 25, 2017
32c15a2
"random op test"
dzhwinter Jul 25, 2017
30a47fe
"link pybind11"
dzhwinter Jul 25, 2017
2b3e362
"template specialization link include"
dzhwinter Jul 25, 2017
984225e
"fix operator"
dzhwinter Jul 25, 2017
11f9f5f
"fix const dependency hell"
dzhwinter Jul 25, 2017
9a16327
"remove const qualify"
dzhwinter Jul 25, 2017
69b1b26
"cpu only macro"
dzhwinter Jul 25, 2017
a22567e
"fix almost equal error"
dzhwinter Jul 25, 2017
5721334
"update the compute kernel"
dzhwinter Jul 30, 2017
36d7e1f
"fix const hell"
dzhwinter Jul 30, 2017
0253f2c
"fix bind python error"
dzhwinter Jul 30, 2017
4d8ece8
"update"
dzhwinter Jul 30, 2017
4755668
"remove unused code"
dzhwinter Jul 30, 2017
4973926
"fix register error"
dzhwinter Jul 30, 2017
933e55e
fix conflict
dzhwinter Jul 30, 2017
2447c34
merge origin/develop
dzhwinter Aug 6, 2017
0f8c9db
device context pointer
dzhwinter Aug 6, 2017
58561d8
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Aug 7, 2017
fcd6f64
"redefine random op"
dzhwinter Aug 7, 2017
e2c08d2
"keep style same with uniform operators"
dzhwinter Aug 8, 2017
52d2ebd
"test gaussian random in python side"
dzhwinter Aug 8, 2017
8804b24
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Aug 8, 2017
555af4d
"format code"
dzhwinter Aug 8, 2017
23ac845
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Aug 8, 2017
6535a7b
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Aug 8, 2017
d98e299
"keep same with uniform random op"
dzhwinter Aug 8, 2017
6fc6647
Merge remote-tracking branch 'origin/develop' into random_op
dzhwinter Aug 8, 2017
7082550
"remove context random seeding "
dzhwinter Aug 8, 2017
df4fe67
"remove attribute"
dzhwinter Aug 9, 2017
6bac3e1
"remove unused test net modified"
dzhwinter Aug 9, 2017
bbd7378
"ci job failed weired. restart ci job."
dzhwinter Aug 9, 2017
f702e79
"relauch ci"
dzhwinter Aug 9, 2017
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
2 changes: 1 addition & 1 deletion paddle/framework/net.h
Original file line number Diff line number Diff line change
Expand Up @@ -71,7 +71,7 @@ class PlainNet : public Net {
* will be used.
*/
void Run(const std::shared_ptr<Scope>& scope,
const platform::DeviceContext& dev_ctx) const override {
platform::DeviceContext& dev_ctx) const override {
for (auto& op : ops_) {
op->Run(scope, dev_ctx);
}
Expand Down
18 changes: 14 additions & 4 deletions paddle/framework/operator.cc
Original file line number Diff line number Diff line change
Expand Up @@ -12,24 +12,24 @@ WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
See the License for the specific language governing permissions and
limitations under the License. */

#include <algorithm>

#include "paddle/framework/operator.h"
#include <algorithm>
#include <iterator>

namespace paddle {
namespace framework {

template <>
Eigen::DefaultDevice* KernelContext::GetEigenDevice<
platform::CPUPlace, Eigen::DefaultDevice>() const {
return device_context_.get_eigen_device<Eigen::DefaultDevice>();
return device_context_->get_eigen_device<Eigen::DefaultDevice>();
}

#ifndef PADDLE_ONLY_CPU
template <>
Eigen::GpuDevice*
KernelContext::GetEigenDevice<platform::GPUPlace, Eigen::GpuDevice>() const {
return device_context_.get_eigen_device<Eigen::GpuDevice>();
return device_context_->get_eigen_device<Eigen::GpuDevice>();
}
#endif

Expand Down Expand Up @@ -95,6 +95,16 @@ std::string OperatorBase::DebugString() const {
ss << ", ";
}
}
ss << "), ";
ss << "Attrs:(";
size_t i = 0;
for (auto& attr : attrs_) {
ss << attr.first;
if (i != attrs_.size() - 1) {
ss << ", ";
}
i++;
}
ss << ").";
return ss.str();
}
Expand Down
8 changes: 4 additions & 4 deletions paddle/framework/operator.h
Original file line number Diff line number Diff line change
Expand Up @@ -114,7 +114,7 @@ class KernelContext {
public:
KernelContext(const OperatorBase* op, const std::shared_ptr<Scope>& scope,
const platform::DeviceContext& device_context)
: op_(*op), scope_(scope), device_context_(device_context) {}
: op_(*op), scope_(scope), device_context_(&device_context) {}

const Variable* Input(int index) const {
return scope_->GetVariable(op_.inputs_[index]);
Expand Down Expand Up @@ -155,11 +155,11 @@ class KernelContext {
typename EigenDeviceConverter<PlaceType>::EigenDeviceType>
DeviceType* GetEigenDevice() const;

platform::Place GetPlace() const { return device_context_.GetPlace(); }
platform::Place GetPlace() const { return device_context_->GetPlace(); }

const OperatorBase& op_;
const std::shared_ptr<Scope>& scope_;
const platform::DeviceContext& device_context_;
const std::shared_ptr<Scope> scope_;
const platform::DeviceContext* device_context_;
};

class OpKernel {
Expand Down
1 change: 1 addition & 0 deletions paddle/operators/CMakeLists.txt
Original file line number Diff line number Diff line change
Expand Up @@ -48,6 +48,7 @@ op_library(mul_op SRCS mul_op.cc mul_op.cu)
op_library(rowwise_add_op SRCS rowwise_add_op.cu rowwise_add_op.cc)
op_library(sigmoid_op SRCS sigmoid_op.cu sigmoid_op.cc)
op_library(softmax_op SRCS softmax_op.cc softmax_op.cu)
op_library(random_op SRCS random_op.cc random_op.cu)
op_library(cross_entropy_op SRCS cross_entropy_op.cc cross_entropy_op.cu)

op_library(fc_op SRCS fc_op.cc DEPS mul_op rowwise_add_op sigmoid_op
Expand Down
86 changes: 86 additions & 0 deletions paddle/operators/random_op.cc
Original file line number Diff line number Diff line change
@@ -0,0 +1,86 @@
/* Copyright (c) 2016 PaddlePaddle Authors. All Rights Reserve.

Licensed under the Apache License, Version 2.0 (the "License");
you may not use this file except in compliance with the License.
You may obtain a copy of the License at

http://www.apache.org/licenses/LICENSE-2.0

Unless required by applicable law or agreed to in writing, software
distributed under the License is distributed on an "AS IS" BASIS,
WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
See the License for the specific language governing permissions and
limitations under the License. */

#include "paddle/operators/random_op.h"
#include "glog/logging.h"
#include "paddle/framework/op_registry.h"

namespace paddle {
namespace operators {

template <typename T>
class GaussianRandomOpKernel<platform::CPUPlace, T>
: public framework::OpKernel {
public:
void Compute(const framework::KernelContext& context) const override {
auto mean = context.op_.GetAttr<T>("mean");
auto std = context.op_.GetAttr<T>("std");
// auto seed = context.op_.GetAttr<T>("seed");
auto* output = context.Output(0)->GetMutable<framework::Tensor>();
T* r = output->mutable_data<T>(context.GetPlace());
auto ctx =
static_cast<const platform::CPUDeviceContext*>(context.device_context_);
// generator need to modify context
auto g = const_cast<platform::CPUDeviceContext*>(ctx)->RandGenerator();
std::normal_distribution<T> distribution(mean, std);
for (int i = 0; i < framework::product(output->dims()); ++i) {
r[i] = distribution(g);
}
}
};

class GaussianRandomOp : public framework::OperatorWithKernel {
protected:
void InferShape(
const std::vector<const framework::Tensor*>& inputs,
const std::vector<framework::Tensor*>& outputs) const override {
PADDLE_ENFORCE(inputs.size() == 0, "Input size of RandomOp must be zero.");
PADDLE_ENFORCE(outputs.size() == 1, "Output size of RandomOp must be one.");
PADDLE_ENFORCE(outputs[0] != nullptr,
"Outputs of RandomOp must all be set.");
outputs[0]->Resize(
framework::make_ddim(this->GetAttr<std::vector<int>>("shape")));
}
};

class GaussianRandomOpMaker : public framework::OpProtoAndCheckerMaker {
public:
GaussianRandomOpMaker(framework::OpProto* proto,
framework::OpAttrChecker* op_checker)
: framework::OpProtoAndCheckerMaker(proto, op_checker) {
AddAttr<std::vector<int>>("shape", "The shape of matrix to be randomized");
// AddAttr<float>("seed", "random seed generator.").SetDefault(1337);
AddAttr<float>("mean", "mean value of random.").SetDefault(.0);
AddAttr<float>("std", "minimum value of random value")
.SetDefault(1.0)
.LargerThan(.0);
AddOutput("Out", "output matrix of random op");
AddComment(R"DOC(
GaussianRandom Operator fill a matrix in normal distribution.
The eqution : Out = GaussianRandom(Shape=(d0, d1, ...), Dtype, mean, std)
)DOC");
}
};

} // namespace operators
} // namespace paddle

REGISTER_OP(gaussian_random,
paddle::operators::GaussianRandomOp,
paddle::operators::GaussianRandomOpMaker);

typedef paddle::operators::GaussianRandomOpKernel<paddle::platform::CPUPlace,
float>
GaussianRandomOpKernel_CPU_float;
REGISTER_OP_CPU_KERNEL(gaussian_random, GaussianRandomOpKernel_CPU_float);
30 changes: 30 additions & 0 deletions paddle/operators/random_op.cu
Original file line number Diff line number Diff line change
@@ -0,0 +1,30 @@
#include "paddle/operators/random_op.h"
#include "paddle/framework/op_registry.h"

namespace paddle {
namespace operators {

template<typename T>
class GaussianRandomOpKernel<platform::GPUPlace, T> : public framework::OpKernel {
public:
void Compute(const framework::KernelContext& context) const override {
auto mean = context.op_.GetAttr<T>("mean");
auto std = context.op_.GetAttr<T>("std");
auto* output = context.Output(0)->GetMutable<framework::Tensor>();
T* r = output->mutable_data<T>(context.GetPlace());
auto ctx = static_cast<const platform::GPUDeviceContext*>
(context.device_context_);
// generator need to modify context
auto g = const_cast<platform::GPUDeviceContext*>(ctx)->RandGenerator();
curandGenerateNormal(g, r, framework::product(output->dims()), mean, std);

}
};

} // namespace operators
} // namespace paddle


typedef paddle::operators::GaussianRandomOpKernel<paddle::platform::GPUPlace, float>
RandomOpKernel_GPU_float;
REGISTER_OP_GPU_KERNEL(random, RandomOpKernel_GPU_float);
17 changes: 17 additions & 0 deletions paddle/operators/random_op.h
Original file line number Diff line number Diff line change
@@ -0,0 +1,17 @@
#pragma once
#include <random>
#include "glog/logging.h"
#include "paddle/framework/eigen.h"
#include "paddle/framework/operator.h"

namespace paddle {
namespace operators {

template <typename Place, typename T>
class GaussianRandomOpKernel : public framework::OpKernel {
public:
void Compute(const framework::KernelContext& context) const override {}
};

} // namespace operators
} // namespace paddle
55 changes: 36 additions & 19 deletions paddle/platform/device_context.h
Original file line number Diff line number Diff line change
Expand Up @@ -21,6 +21,7 @@ limitations under the License. */
#include "paddle/platform/gpu_info.h"
#define EIGEN_USE_GPU
#endif
#include <chrono>
#include <memory>
#include "paddle/platform/place.h"
#include "unsupported/Eigen/CXX11/Tensor"
Expand All @@ -39,7 +40,11 @@ class DeviceContext {

class CPUDeviceContext : public DeviceContext {
public:
CPUDeviceContext() { eigen_device_.reset(new Eigen::DefaultDevice()); }
typedef std::mt19937 random_generator_type;
CPUDeviceContext() {
random_seed_ = std::chrono::system_clock::now().time_since_epoch().count();
Copy link
Collaborator

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

So now we are not able to manually specify random seed?

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

deleted the seeding function in context.
it is the place that we need to modify the Execution context. currently, move the set seed to random op. We need a global configuration to unify the random seed in the future. Maybe we can fix it in next PR.

eigen_device_.reset(new Eigen::DefaultDevice());
}

Eigen::DefaultDevice* eigen_device() const { return eigen_device_.get(); }

Expand All @@ -48,7 +53,16 @@ class CPUDeviceContext : public DeviceContext {
return retv;
}

random_generator_type& RandGenerator() {
if (!rand_generator_) {
rand_generator_.reset(new random_generator_type(random_seed_));
}
return *rand_generator_.get();
}

private:
unsigned random_seed_;
std::unique_ptr<random_generator_type> rand_generator_;
std::unique_ptr<Eigen::DefaultDevice> eigen_device_;
};

Expand All @@ -70,6 +84,9 @@ class GPUPlaceGuard {

class CUDADeviceContext : public DeviceContext {
public:
CUDADeviceContext() {
random_seed_ = std::chrono::system_clock::now().time_since_epoch().count();
}
explicit CUDADeviceContext(const GPUPlace gpu_place) : gpu_place_(gpu_place) {
GPUPlaceGuard guard(gpu_place_);
PADDLE_ENFORCE(cudaStreamCreate(&stream_), "cudaStreamCreate failed");
Expand All @@ -87,6 +104,23 @@ class CUDADeviceContext : public DeviceContext {
"cudaStreamSynchronize failed");
}

curandGenerator_t RandGenerator() {
if (!rand_generator_) {
GPUPlaceGuard guard(gpu_place_);
PADDLE_ENFORCE(paddle::platform::dynload::curandCreateGenerator(
&rand_generator_, CURAND_RNG_PSEUDO_DEFAULT),
"curandCreateGenerator failed");
PADDLE_ENFORCE(
paddle::platform::dynload::curandSetPseudoRandomGeneratorSeed(
rand_generator_, random_seed_),
"curandSetPseudoRandomGeneratorSeed failed");
PADDLE_ENFORCE(
paddle::platform::dynload::curandSetStream(rand_generator_, stream_),
"curandSetStream failed");
}
return rand_generator_;
}

cudaStream_t stream() { return stream_; }

Eigen::GpuDevice* eigen_device() const { return eigen_device_.get(); }
Expand Down Expand Up @@ -115,23 +149,6 @@ class CUDADeviceContext : public DeviceContext {
return dnn_handle_;
}

curandGenerator_t curand_generator() {
if (!rand_generator_) {
GPUPlaceGuard guard(gpu_place_);
PADDLE_ENFORCE(paddle::platform::dynload::curandCreateGenerator(
&rand_generator_, CURAND_RNG_PSEUDO_DEFAULT),
"curandCreateGenerator failed");
PADDLE_ENFORCE(
paddle::platform::dynload::curandSetPseudoRandomGeneratorSeed(
rand_generator_, random_seed_),
"curandSetPseudoRandomGeneratorSeed failed");
PADDLE_ENFORCE(
paddle::platform::dynload::curandSetStream(rand_generator_, stream_),
"curandSetStream failed");
}
return rand_generator_;
}

~CUDADeviceContext() {
Wait();
if (blas_handle_) {
Expand Down Expand Up @@ -165,7 +182,7 @@ class CUDADeviceContext : public DeviceContext {

cudnnHandle_t dnn_handle_{nullptr};

int random_seed_;
unsigned random_seed_;
curandGenerator_t rand_generator_{nullptr};
};

Expand Down
2 changes: 1 addition & 1 deletion paddle/pybind/CMakeLists.txt
Original file line number Diff line number Diff line change
@@ -1,2 +1,2 @@
cc_library(paddle_pybind SHARED SRCS pybind.cc DEPS pybind python
add_op fc_op sgd_op cross_entropy_op)
add_op fc_op sgd_op cross_entropy_op random_op)
1 change: 1 addition & 0 deletions paddle/pybind/pybind.cc
Original file line number Diff line number Diff line change
Expand Up @@ -36,6 +36,7 @@ USE_OP(mul);
USE_OP(sigmoid);
USE_OP(softmax);
USE_OP(rowwise_add);
USE_OP(random);

template <typename ClassType>
void ExposeOperator(ClassType& m) {
Expand Down
3 changes: 2 additions & 1 deletion python/paddle/v2/framework/tests/CMakeLists.txt
Original file line number Diff line number Diff line change
Expand Up @@ -12,4 +12,5 @@ add_python_test(test_framework
test_mul_op.py
test_sigmoid_op.py
test_softmax_op.py
test_rowwise_add_op.py)
test_rowwise_add_op.py
test_random_op.py)
12 changes: 6 additions & 6 deletions python/paddle/v2/framework/tests/test_plain_net.py
Original file line number Diff line number Diff line change
Expand Up @@ -16,13 +16,13 @@ def test_net_all(self):
net.complete_add_op(True)

expected = '''
Op(plain_net), inputs:(@EMPTY@, X, Y, w), outputs:(@TEMP@fc@0, Out, fc.out).
Op(add_two), inputs:(X, Y), outputs:(Out).
Op(plain_net), inputs:(@EMPTY@, X, w), outputs:(@TEMP@fc@0, fc.out).
Op(plain_net), inputs:(@EMPTY@, X, Y, w), outputs:(@TEMP@fc@0, Out, fc.out).
Op(add_two), inputs:(X, Y), outputs:(Out).
Op(plain_net), inputs:(@EMPTY@, X, w), outputs:(@TEMP@fc@0, fc.out).
Op(fc), inputs:(X, w, @EMPTY@), outputs:(fc.out, @TEMP@fc@0).
Op(mul), inputs:(X, w), outputs:(@TEMP@fc@0).
Op(sigmoid), inputs:(@TEMP@fc@0), outputs:(fc.out).
'''
Op(mul), inputs:(X, w), outputs:(@TEMP@fc@0).
Op(sigmoid), inputs:(@TEMP@fc@0), outputs:(fc.out).
'''
self.assertEqual(expected, "\n" + str(net))


Expand Down
29 changes: 29 additions & 0 deletions python/paddle/v2/framework/tests/test_random_op.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,29 @@
import unittest
import paddle.v2.framework.create_op_creation_methods as creation
import paddle.v2.framework.core as core
from op_test_util import OpTestMeta
import numpy


class TestRandomOp(unittest.TestCase):
def test_random(self):
scope = core.Scope(None)
# Out = scope.create_var("Out")
op = creation.op_creations.random(
shape=[1000, 1000], mean=5.0, std=1.0, seed=1701, Out="Out")
for out in op.outputs():
if scope.get_var(out) is None:
scope.create_var(out).get_tensor()

tensor = scope.get_var("Out").get_tensor()
op.infer_shape(scope)
self.assertEqual([1000, 1000], tensor.shape())
ctx = core.DeviceContext.cpu_context()
op.run(scope, ctx)
tensor_array = numpy.array(tensor)
self.assertAlmostEqual(numpy.mean(tensor_array), 5.0, places=3)
self.assertAlmostEqual(numpy.std(tensor_array), 1.0, places=3)


if __name__ == '__main__':
unittest.main()