Fix error message of multinomial op (#27946)

* fix multinomial doc

* fix multinomial error message

* little doc change

* fix Categorical class doc

* optimize format of error message

* fix CPU Kernel error message format

* fix isinf and isnan error in WindowsOPENBLAS CI

* delete inf and nan

* add manual_seed in sample code

* little error message change

* change error message to InvalidArgument

* add full point for error message and add manual_seed in CPU environment
swt-req
pangyoki 4 years ago committed by GitHub
parent b6eff4427c
commit 975bd8873b
No known key found for this signature in database
GPG Key ID: 4AEE18F83AFDEB23

@ -53,12 +53,27 @@ class MultinomialOp : public framework::OperatorWithKernel {
auto x_dim = ctx->GetInputDim("X"); auto x_dim = ctx->GetInputDim("X");
int64_t x_rank = x_dim.size(); int64_t x_rank = x_dim.size();
PADDLE_ENFORCE_GT(x_rank, 0,
platform::errors::InvalidArgument(
"The number of dimensions of the input probability "
"distribution should be > 0, but got %d.",
x_rank));
PADDLE_ENFORCE_LE(x_rank, 2,
platform::errors::InvalidArgument(
"The number of dimensions of the input probability "
"distribution should be <= 2, but got %d.",
x_rank));
std::vector<int64_t> out_dims(x_rank); std::vector<int64_t> out_dims(x_rank);
for (int64_t i = 0; i < x_rank - 1; i++) { for (int64_t i = 0; i < x_rank - 1; i++) {
out_dims[i] = x_dim[i]; out_dims[i] = x_dim[i];
} }
int64_t num_samples = ctx->Attrs().Get<int>("num_samples"); int64_t num_samples = ctx->Attrs().Get<int>("num_samples");
PADDLE_ENFORCE_GT(
num_samples, 0,
platform::errors::InvalidArgument(
"The number of samples should be > 0, but got %d.", num_samples));
out_dims[x_rank - 1] = num_samples; out_dims[x_rank - 1] = num_samples;
ctx->SetOutputDim("Out", framework::make_ddim(out_dims)); ctx->SetOutputDim("Out", framework::make_ddim(out_dims));

@ -21,6 +21,7 @@ limitations under the License. */
#include "paddle/fluid/framework/op_registry.h" #include "paddle/fluid/framework/op_registry.h"
#include "paddle/fluid/framework/operator.h" #include "paddle/fluid/framework/operator.h"
#include "paddle/fluid/operators/multinomial_op.h" #include "paddle/fluid/operators/multinomial_op.h"
#include "paddle/fluid/platform/enforce.h"
#include "paddle/fluid/platform/transform.h" #include "paddle/fluid/platform/transform.h"
namespace paddle { namespace paddle {
@ -31,6 +32,14 @@ __global__ void NormalizeProbability(T* norm_probs, const T* in_data,
T* sum_rows) { T* sum_rows) {
int id = threadIdx.x + blockIdx.x * blockDim.x + int id = threadIdx.x + blockIdx.x * blockDim.x +
blockIdx.y * gridDim.x * blockDim.x; blockIdx.y * gridDim.x * blockDim.x;
PADDLE_ENFORCE(
in_data[id] >= 0.0,
"The input of multinomial distribution should be >= 0, but got %f.",
in_data[id]);
PADDLE_ENFORCE(sum_rows[blockIdx.y] > 0.0,
"The sum of one multinomial distribution probability should "
"be > 0, but got %f.",
sum_rows[blockIdx.y]);
norm_probs[id] = in_data[id] / sum_rows[blockIdx.y]; norm_probs[id] = in_data[id] / sum_rows[blockIdx.y];
} }

@ -44,28 +44,29 @@ void MultinomialFunctor(int64_t* out_data, const T* in_data,
int64_t num_zeros = 0; int64_t num_zeros = 0;
for (int64_t j = 0; j < num_categories; j++) { for (int64_t j = 0; j < num_categories; j++) {
prob_value = in_data[i * num_categories + j]; prob_value = in_data[i * num_categories + j];
PADDLE_ENFORCE_GE( PADDLE_ENFORCE_GE(prob_value, 0.0,
prob_value, 0.0, platform::errors::InvalidArgument(
platform::errors::OutOfRange( "The input of multinomial distribution "
"The input of multinomial distribution should be >= 0")); "should be >= 0, but got %f.",
PADDLE_ENFORCE_EQ((std::isinf(static_cast<double>(prob_value)) || prob_value));
std::isnan(static_cast<double>(prob_value))),
false, platform::errors::OutOfRange(
"The input of multinomial distribution "
"shoud not be infinity or NaN"));
probs_sum += prob_value; probs_sum += prob_value;
if (prob_value == 0) { if (prob_value == 0) {
num_zeros += 1; num_zeros += 1;
} }
cumulative_probs[j] = probs_sum; cumulative_probs[j] = probs_sum;
} }
PADDLE_ENFORCE_GT(probs_sum, 0.0, platform::errors::OutOfRange( PADDLE_ENFORCE_GT(probs_sum, 0.0,
"The sum of input should not be 0")); platform::errors::InvalidArgument(
"The sum of one multinomial distribution "
"probability should be > 0, but got %f.",
probs_sum));
PADDLE_ENFORCE_EQ( PADDLE_ENFORCE_EQ(
(replacement || (num_categories - num_zeros >= num_samples)), true, (replacement || (num_categories - num_zeros >= num_samples)), true,
platform::errors::OutOfRange("When replacement is False, number of " platform::errors::InvalidArgument(
"samples should be less than non-zero " "When replacement is False, number of "
"categories")); "samples should be less than non-zero "
"categories."));
for (int64_t j = 0; j < num_categories; j++) { for (int64_t j = 0; j < num_categories; j++) {
cumulative_probs[j] /= probs_sum; cumulative_probs[j] /= probs_sum;

File diff suppressed because it is too large Load Diff

@ -17,12 +17,14 @@ from __future__ import print_function
import unittest import unittest
import paddle import paddle
import paddle.fluid as fluid import paddle.fluid as fluid
from paddle.fluid import core
from op_test import OpTest from op_test import OpTest
import numpy as np import numpy as np
class TestMultinomialOp(OpTest): class TestMultinomialOp(OpTest):
def setUp(self): def setUp(self):
paddle.enable_static()
self.op_type = "multinomial" self.op_type = "multinomial"
self.init_data() self.init_data()
self.inputs = {"X": self.input_np} self.inputs = {"X": self.input_np}
@ -175,5 +177,39 @@ class TestMultinomialAlias(unittest.TestCase):
paddle.tensor.random.multinomial(x, num_samples=10, replacement=True) paddle.tensor.random.multinomial(x, num_samples=10, replacement=True)
class TestMultinomialError(unittest.TestCase):
def setUp(self):
paddle.disable_static()
def test_num_sample(self):
def test_num_sample_less_than_0():
x = paddle.rand([4])
paddle.multinomial(x, num_samples=-2)
self.assertRaises(core.EnforceNotMet, test_num_sample_less_than_0)
def test_replacement_False(self):
def test_samples_larger_than_categories():
x = paddle.rand([4])
paddle.multinomial(x, num_samples=5, replacement=False)
self.assertRaises(core.EnforceNotMet,
test_samples_larger_than_categories)
def test_input_probs_dim(self):
def test_dim_larger_than_2():
x = paddle.rand([2, 3, 3])
paddle.multinomial(x)
self.assertRaises(core.EnforceNotMet, test_dim_larger_than_2)
def test_dim_less_than_1():
x_np = np.random.random([])
x = paddle.to_tensor(x_np)
paddle.multinomial(x)
self.assertRaises(core.EnforceNotMet, test_dim_less_than_1)
if __name__ == "__main__": if __name__ == "__main__":
unittest.main() unittest.main()

@ -57,19 +57,19 @@ def bernoulli(x, name=None):
Examples: Examples:
.. code-block:: python .. code-block:: python
import paddle import paddle
paddle.disable_static()
x = paddle.rand([2, 3]) paddle.manual_seed(100) # on CPU device
print(x.numpy()) x = paddle.rand([2,3])
# [[0.11272584 0.3890902 0.7730957 ] print(x.numpy())
# [0.10351662 0.8510418 0.63806665]] # [[0.5535528 0.20714243 0.01162981]
# [0.51577556 0.36369765 0.2609165 ]]
out = paddle.bernoulli(x) paddle.manual_seed(200) # on CPU device
print(out.numpy()) out = paddle.bernoulli(x)
# [[0. 0. 1.] print(out.numpy())
# [0. 0. 1.]] # [[0. 0. 0.]
# [1. 1. 0.]]
""" """
@ -108,28 +108,29 @@ def multinomial(x, num_samples=1, replacement=False, name=None):
Examples: Examples:
.. code-block:: python .. code-block:: python
import paddle import paddle
paddle.disable_static()
x = paddle.rand([2,4])
print(x.numpy())
# [[0.7713825 0.4055941 0.433339 0.70706886]
# [0.9223313 0.8519825 0.04574518 0.16560672]]
out1 = paddle.multinomial(x, num_samples=5, replacement=True)
print(out1.numpy())
# [[3 3 1 1 0]
# [0 0 0 0 1]]
# out2 = paddle.multinomial(x, num_samples=5)
# OutOfRangeError: When replacement is False, number of samples
# should be less than non-zero categories
out3 = paddle.multinomial(x, num_samples=3) paddle.manual_seed(100) # on CPU device
print(out3.numpy()) x = paddle.rand([2,4])
# [[0 2 3] print(x.numpy())
# [0 1 3]] # [[0.5535528 0.20714243 0.01162981 0.51577556]
# [0.36369765 0.2609165 0.18905126 0.5621971 ]]
paddle.manual_seed(200) # on CPU device
out1 = paddle.multinomial(x, num_samples=5, replacement=True)
print(out1.numpy())
# [[3 3 0 0 0]
# [3 3 3 1 0]]
# out2 = paddle.multinomial(x, num_samples=5)
# InvalidArgumentError: When replacement is False, number of samples
# should be less than non-zero categories
paddle.manual_seed(300) # on CPU device
out3 = paddle.multinomial(x, num_samples=3)
print(out3.numpy())
# [[3 0 1]
# [3 1 0]]
""" """

Loading…
Cancel
Save